llm_vllm.py 11 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234
  1. # Copyright (c) 2024 Alibaba Inc (authors: Xiang Lyu, Zhihao Du)
  2. #
  3. # Licensed under the Apache License, Version 2.0 (the "License");
  4. # you may not use this file except in compliance with the License.
  5. # You may obtain a copy of the License at
  6. #
  7. # http://www.apache.org/licenses/LICENSE-2.0
  8. #
  9. # Unless required by applicable law or agreed to in writing, software
  10. # distributed under the License is distributed on an "AS IS" BASIS,
  11. # WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
  12. # See the License for the specific language governing permissions and
  13. # limitations under the License.
  14. import time
  15. import queue
  16. import asyncio
  17. import threading
  18. from typing import List, Generator, AsyncGenerator
  19. import torch
  20. from cosyvoice.utils.file_utils import logging
  21. from cosyvoice.llm.llm import Qwen2LM
  22. # 启用vllm V1版本
  23. import os
  24. os.environ["VLLM_USE_V1"] = '1'
  25. from vllm import ModelRegistry
  26. from vllm import LLMEngine, AsyncLLMEngine, CompletionOutput
  27. from vllm.engine.arg_utils import EngineArgs, AsyncEngineArgs
  28. from vllm.sampling_params import SamplingParams
  29. from cosyvoice.llm.vllm_use_cosyvoice2_model import CosyVoice2Model as CosyVoice2LLM
  30. ModelRegistry.register_model("CosyVoice2Model", CosyVoice2LLM)
  31. # EngineArgs
  32. ENGINE_ARGS = {
  33. "block_size": 16,
  34. "swap_space": 0,
  35. # "enforce_eager": True,
  36. "gpu_memory_utilization": 0.4,
  37. "max_num_batched_tokens": 1024,
  38. "max_model_len": 1024,
  39. "max_num_seqs": 256,
  40. "disable_log_requests": True,
  41. "disable_log_stats": True,
  42. "dtype": "float16"
  43. }
  44. from vllm.sampling_params import RequestOutputKind
  45. # SamplingParams
  46. SAMPLING_PARAMS = {
  47. "temperature": 1, # 不能低于0.8, 否则会生成非常多的空音频,或者无法正常生成语音Token
  48. "top_p": 1, # 不能低于0.8, 否则会生成非常多的空音频,或者无法正常生成语音Token
  49. "top_k": 25,
  50. # "min_tokens": 80, # 不支持设置最小的tokens数量设置,开启后vllm直接崩溃,无法启动
  51. # "presence_penalty": 1.0, # 不支持设置
  52. # "frequency_penalty": 0.0, # 不支持设置
  53. "max_tokens": 1024,
  54. "detokenize": False, # 目前 vllm 0.7.3 v1版本中设置无效,待后续版本更新后减少计算
  55. "ignore_eos": False,
  56. "output_kind": RequestOutputKind.DELTA # 设置为DELTA,如调整该参数,请同时调整llm_inference的处理代码
  57. }
  58. def tensor_to_list(tensor: torch.tensor):
  59. return tensor.view(-1).cpu().numpy().tolist()
  60. class VllmQwen2LM(Qwen2LM):
  61. def __init__(
  62. self,
  63. model_dir,
  64. mix_ratio: List[int] = [5, 15],
  65. ):
  66. self.fp16 = False
  67. self.half = lambda: None
  68. self.mix_ratio = mix_ratio
  69. # ---------------------------------------------
  70. # vllm engine 的参数配置
  71. engine_args = AsyncEngineArgs(
  72. model=model_dir,
  73. **ENGINE_ARGS,
  74. )
  75. self.llm_engine: AsyncLLMEngine = AsyncLLMEngine.from_engine_args(engine_args)
  76. self.speech_token_size = 6564 # 6561 + 3
  77. self.llm_token_size = 151936 # llm vocab_size
  78. self.sos_eos_token_id = self.speech_token_size + self.llm_token_size + 1
  79. self.task_token_id = self.sos_eos_token_id + 1
  80. self.zero_token_id = self.task_token_id + 1
  81. # 不能直接在同步函数正确的使用 异步的生成器函数,即使使用协程也会对vllm造成崩溃
  82. # 使用 queue 的方式,后台线程运行推理任务
  83. self.task_queue = queue.Queue()
  84. self.loop = asyncio.new_event_loop()
  85. self.loop_thread = threading.Thread(target=self._run_event_loop, daemon=True)
  86. self.loop_thread.start()
  87. # 运行后台协程,用于处理任务队列中的任务
  88. # TODO: 目前只能单任务运行,多任务运行需要对 inference_processor 进行修改
  89. asyncio.run_coroutine_threadsafe(self.inference_processor(self.task_queue), self.loop)
  90. def _run_event_loop(self):
  91. asyncio.set_event_loop(self.loop)
  92. self.loop.run_forever()
  93. async def inference_processor(self, task_queue):
  94. while True:
  95. try:
  96. print(f"inference_processor")
  97. out_queue, prompt_token_ids, request_id, stop_token_ids, max_tokens = task_queue.get()
  98. sampling_params = SamplingParams(**SAMPLING_PARAMS)
  99. sampling_params.stop_token_ids = stop_token_ids or [6561]
  100. if max_tokens:
  101. sampling_params.max_tokens = max_tokens
  102. async for output in self.llm_engine.generate(
  103. {
  104. "prompt_token_ids": prompt_token_ids,
  105. },
  106. sampling_params=sampling_params,
  107. request_id=request_id or f"{time.time()}",
  108. ):
  109. out_queue.put((output.outputs[0], output.finished))
  110. except Exception as e:
  111. logging.error(f"Error in inference_processor: {e}")
  112. def llm_inference(self, prompt_token_ids: List[int], request_id: str=None, stop_token_ids=None, max_tokens=None):
  113. # 使用 同步转异步 会导致vllm崩溃,目前选择 queue 的方式,后台线程运行推理任务
  114. # 提交推理任务到队列中
  115. out_queue = queue.Queue()
  116. self.task_queue.put((out_queue, prompt_token_ids, request_id, stop_token_ids, max_tokens))
  117. # 将 out_queue 的结果返回
  118. finished = False
  119. while not finished:
  120. (output, finished) = out_queue.get_nowait() if not out_queue.empty() else out_queue.get()
  121. yield output
  122. def inference(
  123. self,
  124. text: torch.Tensor,
  125. text_len: torch.Tensor,
  126. prompt_text: torch.Tensor,
  127. prompt_text_len: torch.Tensor,
  128. prompt_speech_token: torch.Tensor,
  129. prompt_speech_token_len: torch.Tensor,
  130. embedding: torch.Tensor,
  131. sampling: int = 25,
  132. max_token_text_ratio: float = 20,
  133. min_token_text_ratio: float = 2,
  134. ) -> Generator[torch.Tensor|int, None, None]:
  135. prompt_text = tensor_to_list(prompt_text + torch.tensor(6564))
  136. prompt_speech_token = tensor_to_list(prompt_speech_token)
  137. text = tensor_to_list(text + torch.tensor(6564))
  138. prompt_token_ids = [self.sos_eos_token_id] + prompt_text + text + \
  139. [self.task_token_id] + prompt_speech_token
  140. max_tokens = len(text) * 20
  141. for output in self.llm_inference(
  142. prompt_token_ids,
  143. stop_token_ids=[6561],
  144. max_tokens=max_tokens,
  145. ):
  146. if output.token_ids[-1] == 6561:
  147. need_add_tokens = output.token_ids[:-1]
  148. else:
  149. need_add_tokens = output.token_ids
  150. # 单个token 循环处理比较耗时,建议是在model中进行批量(extend)处理,减少循环
  151. # yield need_add_tokens
  152. for token in need_add_tokens:
  153. yield token
  154. def inference_bistream(
  155. self,
  156. text: Generator,
  157. prompt_text: torch.Tensor,
  158. prompt_text_len: torch.Tensor,
  159. prompt_speech_token: torch.Tensor,
  160. prompt_speech_token_len: torch.Tensor,
  161. embedding: torch.Tensor,
  162. sampling: int = 25,
  163. max_token_text_ratio: float = 20,
  164. min_token_text_ratio: float = 2,
  165. ) -> Generator[torch.Tensor, None, None]:
  166. prompt_text = tensor_to_list(prompt_text + torch.tensor(6564))
  167. prompt_speech_token = tensor_to_list(prompt_speech_token)
  168. last_tokens = []
  169. prompt_token_ids = [self.sos_eos_token_id]
  170. text_tokens_cache = prompt_text
  171. for this_text in text:
  172. this_text = tensor_to_list(this_text + torch.tensor(6564))
  173. # text need tokens
  174. assert isinstance(this_text, list), "text need token ids List[int]."
  175. text_tokens_cache += this_text
  176. while len(prompt_speech_token) != 0:
  177. if len(text_tokens_cache) >= self.mix_ratio[0]:
  178. text_input_token = text_tokens_cache[:self.mix_ratio[0]]
  179. speech_input_token = prompt_speech_token[:self.mix_ratio[1]]
  180. prompt_token_ids += text_input_token + speech_input_token
  181. # reset the last cache
  182. text_tokens_cache = text_tokens_cache[self.mix_ratio[0]:]
  183. prompt_speech_token = prompt_speech_token[self.mix_ratio[1]:]
  184. else:
  185. logging.info('not enough text token to decode, wait for more')
  186. break
  187. if len(prompt_speech_token) == 0:
  188. if (len(last_tokens) > 0 and last_tokens[-1] == 6563) or len(prompt_token_ids) == 1:
  189. logging.info('get fill token, need to append more text token')
  190. if len(text_tokens_cache) >= self.mix_ratio[0]:
  191. text_tokens_temp = text_tokens_cache[:self.mix_ratio[0]]
  192. prompt_token_ids += text_tokens_temp
  193. logging.info('append {} text token'.format(len(text_tokens_temp)))
  194. text_tokens_cache = text_tokens_cache[self.mix_ratio[0]:]
  195. else:
  196. logging.info('not enough text token to decode, wait for more')
  197. continue
  198. for output in self.llm_inference(prompt_token_ids, stop_token_ids=[6563]):
  199. last_tokens = output.token_ids
  200. if last_tokens[-1] == 6563:
  201. need_add_tokens = last_tokens[:-1]
  202. else:
  203. need_add_tokens = last_tokens
  204. # 单个token 循环处理比较耗时,建议是在model中进行批量(extend)处理,减少循环
  205. # yield need_add_tokens
  206. for token in need_add_tokens:
  207. yield token
  208. prompt_token_ids.extend(need_add_tokens)
  209. prompt_token_ids += text_tokens_cache + [self.task_token_id]
  210. logging.info('no more text token, decode until met eos')
  211. for output in self.llm_inference(prompt_token_ids, stop_token_ids=[6561]):
  212. if output.token_ids[-1] == 6561:
  213. need_add_tokens = output.token_ids[:-1]
  214. else:
  215. need_add_tokens = output.token_ids
  216. # 单个token 循环处理比较耗时,建议是在model中进行批量(extend)处理,减少循环
  217. # yield need_add_tokens
  218. for token in need_add_tokens:
  219. yield token