worker.py 17 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394
  1. """A GPU worker class."""
  2. import gc
  3. import os
  4. from typing import List, Optional, Set, Tuple, Type
  5. import torch
  6. import torch.distributed
  7. from loguru import logger
  8. from aphrodite.common.config import (CacheConfig, DeviceConfig, LoadConfig,
  9. LoRAConfig, ModelConfig, MultiModalConfig,
  10. ParallelConfig, PromptAdapterConfig,
  11. SchedulerConfig, SpeculativeConfig)
  12. from aphrodite.common.sequence import ExecuteModelRequest
  13. from aphrodite.distributed import (ensure_model_parallel_initialized,
  14. init_distributed_environment,
  15. set_custom_all_reduce)
  16. from aphrodite.lora.request import LoRARequest
  17. from aphrodite.modeling import set_random_seed
  18. from aphrodite.modeling.model_loader.tensorizer import TensorizerConfig
  19. from aphrodite.platforms import current_platform
  20. from aphrodite.prompt_adapter.request import PromptAdapterRequest
  21. from aphrodite.task_handler.cache_engine import CacheEngine
  22. from aphrodite.task_handler.embedding_model_runner import EmbeddingModelRunner
  23. from aphrodite.task_handler.model_runner import GPUModelRunnerBase, ModelRunner
  24. from aphrodite.task_handler.worker_base import (LocalOrDistributedWorkerBase,
  25. WorkerInput)
  26. class Worker(LocalOrDistributedWorkerBase):
  27. """A worker class that executes (a partition of) the model on a GPU.
  28. Each worker is associated with a single GPU. The worker is responsible for
  29. maintaining the KV cache and executing the model on the GPU. In case of
  30. distributed inference, each worker is assigned a partition of the model.
  31. """
  32. def __init__(
  33. self,
  34. model_config: ModelConfig,
  35. parallel_config: ParallelConfig,
  36. scheduler_config: SchedulerConfig,
  37. device_config: DeviceConfig,
  38. cache_config: CacheConfig,
  39. load_config: LoadConfig,
  40. local_rank: int,
  41. rank: int,
  42. distributed_init_method: str,
  43. lora_config: Optional[LoRAConfig] = None,
  44. multimodal_config: Optional[MultiModalConfig] = None,
  45. speculative_config: Optional[SpeculativeConfig] = None,
  46. prompt_adapter_config: Optional[PromptAdapterConfig] = None,
  47. is_driver_worker: bool = False,
  48. model_runner_cls: Optional[Type[GPUModelRunnerBase]] = None,
  49. ) -> None:
  50. self.model_config = model_config
  51. self.parallel_config = parallel_config
  52. self.parallel_config.rank = rank
  53. self.scheduler_config = scheduler_config
  54. self.device_config = device_config
  55. self.cache_config = cache_config
  56. self.local_rank = local_rank
  57. self.rank = rank
  58. self.distributed_init_method = distributed_init_method
  59. self.lora_config = lora_config
  60. self.prompt_adapter_config = prompt_adapter_config
  61. self.load_config = load_config
  62. self.is_driver_worker = is_driver_worker
  63. if parallel_config and is_driver_worker:
  64. assert rank % parallel_config.tensor_parallel_size == 0, \
  65. "Driver worker should be rank 0 of tensor parallel group."
  66. if self.model_config.trust_remote_code:
  67. # note: lazy import to avoid importing torch before initializing
  68. from aphrodite.common.utils import init_cached_hf_modules
  69. init_cached_hf_modules()
  70. self.multimodal_config = multimodal_config
  71. # Return hidden states from target model if the draft model is an
  72. # mlp_speculator
  73. speculative_args = {} if speculative_config is None \
  74. or (speculative_config.draft_model_config.model ==
  75. model_config.model) \
  76. or (speculative_config.draft_model_config.hf_config.model_type
  77. not in ["medusa", "mlp_speculator"]) \
  78. else {"return_hidden_states": True}
  79. ModelRunnerClass: Type[GPUModelRunnerBase] = ModelRunner
  80. if model_runner_cls is not None:
  81. ModelRunnerClass = model_runner_cls
  82. elif self.model_config.embedding_mode:
  83. ModelRunnerClass = EmbeddingModelRunner
  84. self.model_runner: GPUModelRunnerBase = ModelRunnerClass(
  85. model_config,
  86. parallel_config,
  87. scheduler_config,
  88. device_config,
  89. cache_config,
  90. load_config=load_config,
  91. lora_config=self.lora_config,
  92. kv_cache_dtype=self.cache_config.cache_dtype,
  93. is_driver_worker=is_driver_worker,
  94. prompt_adapter_config=prompt_adapter_config,
  95. multimodal_config=multimodal_config,
  96. **speculative_args,
  97. )
  98. # Uninitialized cache engine. Will be initialized by
  99. # initialize_cache.
  100. self.cache_engine: List[CacheEngine]
  101. # Initialize gpu_cache as embedding models don't initialize kv_caches
  102. self.gpu_cache: Optional[List[List[torch.tensor]]] = None
  103. def init_device(self) -> None:
  104. if self.device_config.device.type == "cuda":
  105. # torch.distributed.all_reduce does not free the input tensor until
  106. # the synchronization point. This causes the memory usage to grow
  107. # as the number of all_reduce calls increases. This env var disables
  108. # this behavior.
  109. # Related issue:
  110. # https://discuss.pytorch.org/t/cuda-allocation-lifetime-for-inputs-to-distributed-all-reduce/191573
  111. os.environ["TORCH_NCCL_AVOID_RECORD_STREAMS"] = "1"
  112. # This env var set by Ray causes exceptions with graph building.
  113. os.environ.pop("NCCL_ASYNC_ERROR_HANDLING", None)
  114. self.device = torch.device(f"cuda:{self.local_rank}")
  115. torch.cuda.set_device(self.device)
  116. _check_if_gpu_supports_dtype(self.model_config.dtype)
  117. torch.cuda.empty_cache()
  118. self.init_gpu_memory = torch.cuda.mem_get_info()[0]
  119. else:
  120. raise RuntimeError(
  121. f"Not support device type: {self.device_config.device}")
  122. # Initialize the distributed environment.
  123. init_worker_distributed_environment(self.parallel_config, self.rank,
  124. self.distributed_init_method,
  125. self.local_rank)
  126. # Set random seed.
  127. set_random_seed(self.model_config.seed)
  128. def load_model(self):
  129. self.model_runner.load_model()
  130. def save_sharded_state(
  131. self,
  132. path: str,
  133. pattern: Optional[str] = None,
  134. max_size: Optional[int] = None,
  135. ) -> None:
  136. self.model_runner.save_sharded_state(
  137. path,
  138. pattern=pattern,
  139. max_size=max_size,
  140. )
  141. def save_tensorized_model(
  142. self,
  143. tensorizer_config: TensorizerConfig,
  144. ) -> None:
  145. self.model_runner.save_tensorized_model(
  146. tensorizer_config=tensorizer_config, )
  147. @torch.inference_mode()
  148. def determine_num_available_blocks(self) -> Tuple[int, int]:
  149. """Profiles the peak memory usage of the model to determine how many
  150. KV blocks may be allocated without OOMs.
  151. The engine will first conduct a profiling of the existing memory usage.
  152. Then, it calculate the maximum possible number of GPU and CPU blocks
  153. that can be allocated with the remaining free memory.
  154. .. tip::
  155. You may limit the usage of GPU memory
  156. by adjusting the `gpu_memory_utilization` parameter.
  157. """
  158. # Profile the memory usage of the model and get the maximum number of
  159. # cache blocks that can be allocated with the remaining free memory.
  160. torch.cuda.empty_cache()
  161. # Execute a forward pass with dummy inputs to profile the memory usage
  162. # of the model.
  163. self.model_runner.profile_run()
  164. # Calculate the number of blocks that can be allocated with the
  165. # profiled peak memory.
  166. torch.cuda.synchronize()
  167. free_gpu_memory, total_gpu_memory = torch.cuda.mem_get_info()
  168. # NOTE: Here we assume that the other processes using the same
  169. # GPU did not change their memory usage during the profiling.
  170. peak_memory = self.init_gpu_memory - free_gpu_memory
  171. assert peak_memory > 0, (
  172. "Error in memory profiling. This happens when the GPU memory was "
  173. "not properly cleaned up before initializing Aphrodite.")
  174. cache_block_size = self.get_cache_block_size_bytes()
  175. num_gpu_blocks = int(
  176. (total_gpu_memory * self.cache_config.gpu_memory_utilization -
  177. peak_memory) // cache_block_size)
  178. num_cpu_blocks = int(self.cache_config.swap_space_bytes //
  179. cache_block_size)
  180. num_gpu_blocks = max(num_gpu_blocks, 0)
  181. num_cpu_blocks = max(num_cpu_blocks, 0)
  182. if self.model_runner.lora_manager:
  183. self.model_runner.remove_all_loras()
  184. gc.collect()
  185. torch.cuda.empty_cache()
  186. return num_gpu_blocks, num_cpu_blocks
  187. def initialize_cache(self, num_gpu_blocks: int,
  188. num_cpu_blocks: int) -> None:
  189. """Allocate GPU and CPU KV cache with the specified number of blocks.
  190. This also warms up the model, which may record CUDA graphs.
  191. """
  192. raise_if_cache_size_invalid(num_gpu_blocks,
  193. self.cache_config.block_size,
  194. self.model_config.max_model_len)
  195. self.cache_config.num_gpu_blocks = num_gpu_blocks
  196. self.cache_config.num_cpu_blocks = num_cpu_blocks
  197. self._init_cache_engine()
  198. self._warm_up_model()
  199. def _init_cache_engine(self):
  200. assert self.cache_config.num_gpu_blocks is not None
  201. self.cache_engine = [
  202. CacheEngine(self.cache_config, self.model_config,
  203. self.parallel_config, self.device_config)
  204. for _ in range(self.parallel_config.pipeline_parallel_size)
  205. ]
  206. self.gpu_cache = [
  207. self.cache_engine[ve].gpu_cache
  208. for ve in range(self.parallel_config.pipeline_parallel_size)
  209. ]
  210. def _warm_up_model(self) -> None:
  211. if not self.model_config.enforce_eager:
  212. self.model_runner.capture_model(self.gpu_cache)
  213. # Reset the seed to ensure that the random state is not affected by
  214. # the model initialization and profiling.
  215. set_random_seed(self.model_config.seed)
  216. @property
  217. def do_metadata_broadcast(self) -> bool:
  218. return self.parallel_config.tensor_parallel_size > 1
  219. @property
  220. def kv_cache(self) -> Optional[List[List[torch.Tensor]]]:
  221. return self.gpu_cache
  222. @torch.inference_mode()
  223. def prepare_worker_input(
  224. self, execute_model_req: ExecuteModelRequest) -> WorkerInput:
  225. virtual_engine = execute_model_req.virtual_engine
  226. num_seq_groups = len(execute_model_req.seq_group_metadata_list)
  227. # `blocks_to_swap_in` and `blocks_to_swap_out` are cpu tensors.
  228. # they contain parameters to launch cudamemcpyasync.
  229. blocks_to_swap_in = torch.tensor(execute_model_req.blocks_to_swap_in,
  230. device="cpu",
  231. dtype=torch.int64).view(-1, 2)
  232. blocks_to_swap_out = torch.tensor(execute_model_req.blocks_to_swap_out,
  233. device="cpu",
  234. dtype=torch.int64).view(-1, 2)
  235. # `blocks_to_copy` is a gpu tensor. The src and tgt of
  236. # blocks to copy are in the same device, and `blocks_to_copy`
  237. # can be used directly within cuda kernels.
  238. blocks_to_copy = torch.tensor(execute_model_req.blocks_to_copy,
  239. device=self.device,
  240. dtype=torch.int64).view(-1, 2)
  241. return WorkerInput(num_seq_groups=num_seq_groups,
  242. blocks_to_swap_in=blocks_to_swap_in,
  243. blocks_to_swap_out=blocks_to_swap_out,
  244. blocks_to_copy=blocks_to_copy,
  245. virtual_engine=virtual_engine)
  246. @torch.inference_mode()
  247. def execute_worker(self, worker_input: WorkerInput) -> None:
  248. virtual_engine = worker_input.virtual_engine
  249. # Issue cache operations.
  250. if (worker_input.blocks_to_swap_in is not None
  251. and worker_input.blocks_to_swap_in.numel() > 0):
  252. self.cache_engine[virtual_engine].swap_in(
  253. worker_input.blocks_to_swap_in)
  254. if (worker_input.blocks_to_swap_out is not None
  255. and worker_input.blocks_to_swap_out.numel() > 0):
  256. self.cache_engine[virtual_engine].swap_out(
  257. worker_input.blocks_to_swap_out)
  258. if (worker_input.blocks_to_copy is not None
  259. and worker_input.blocks_to_copy.numel() > 0):
  260. self.cache_engine[virtual_engine].copy(worker_input.blocks_to_copy)
  261. def add_lora(self, lora_request: LoRARequest) -> bool:
  262. return self.model_runner.add_lora(lora_request)
  263. def remove_lora(self, lora_id: int) -> bool:
  264. return self.model_runner.remove_lora(lora_id)
  265. def pin_lora(self, lora_id: int) -> bool:
  266. return self.model_runner.pin_lora(lora_id)
  267. def list_loras(self) -> Set[int]:
  268. return self.model_runner.list_loras()
  269. def add_prompt_adapter(
  270. self, prompt_adapter_request: PromptAdapterRequest) -> bool:
  271. return self.model_runner.add_prompt_adapter(prompt_adapter_request)
  272. def remove_prompt_adapter(self, prompt_adapter_id: int) -> bool:
  273. return self.model_runner.remove_lora(prompt_adapter_id)
  274. def pin_prompt_adapter(self, prompt_adapter_id: int) -> bool:
  275. return self.model_runner.pin_prompt_adapter(prompt_adapter_id)
  276. def list_prompt_adapters(self) -> Set[int]:
  277. return self.model_runner.list_prompt_adapters()
  278. @property
  279. def max_model_len(self) -> int:
  280. return self.model_config.max_model_len
  281. @property
  282. def vocab_size(self) -> int:
  283. return self.model_runner.vocab_size
  284. def get_cache_block_size_bytes(self) -> int:
  285. """Get the size of the KV cache block size in bytes.
  286. """
  287. return CacheEngine.get_cache_block_size(self.cache_config,
  288. self.model_config,
  289. self.parallel_config)
  290. def init_worker_distributed_environment(
  291. parallel_config: ParallelConfig,
  292. rank: int,
  293. distributed_init_method: Optional[str] = None,
  294. local_rank: int = -1,
  295. ) -> None:
  296. """Initialize the distributed environment."""
  297. set_custom_all_reduce(not parallel_config.disable_custom_all_reduce)
  298. init_distributed_environment(parallel_config.world_size, rank,
  299. distributed_init_method, local_rank)
  300. ensure_model_parallel_initialized(parallel_config.tensor_parallel_size,
  301. parallel_config.pipeline_parallel_size)
  302. def _check_if_gpu_supports_dtype(torch_dtype: torch.dtype):
  303. # Check if the GPU supports the dtype.
  304. if torch_dtype == torch.bfloat16:
  305. compute_capability = current_platform.get_device_capability()
  306. if compute_capability[0] < 8:
  307. gpu_name = torch.cuda.get_device_name()
  308. raise ValueError(
  309. "Bfloat16 is only supported on GPUs with compute capability "
  310. f"of at least 8.0. Your {gpu_name} GPU has compute capability "
  311. f"{compute_capability[0]}.{compute_capability[1]}. "
  312. "You can use float16 instead by explicitly setting the"
  313. "`dtype` flag in CLI, for example: --dtype=half.")
  314. def raise_if_cache_size_invalid(num_gpu_blocks, block_size,
  315. max_model_len) -> None:
  316. if num_gpu_blocks <= 0:
  317. raise ValueError("No available memory for the cache blocks. "
  318. "Try increasing `gpu_memory_utilization` when "
  319. "initializing the engine.")
  320. max_seq_len = block_size * num_gpu_blocks
  321. logger.info(f"Maximum sequence length allowed in the cache: "
  322. f"{max_seq_len}")
  323. if max_model_len > max_seq_len:
  324. original_max_model_len = max_model_len
  325. max_model_len = max_seq_len
  326. # raise ValueError(
  327. # f"The model's max seq len ({max_model_len}) "
  328. # "is larger than the maximum number of tokens that can be "
  329. # f"stored in KV cache ({max_seq_len}). Try increasing "
  330. # "`gpu_memory_utilization` or decreasing `max_model_len` when "
  331. # "initializing the engine.")
  332. # set the max_model_len to the max_seq_len, but raise a logger.error
  333. # so the user is made aware of this
  334. logger.error(
  335. f"The model's max seq len ({original_max_model_len}) "
  336. "is larger than the maximum number of tokens that can be "
  337. f"stored in KV cache ({max_seq_len}). "
  338. "Try increasing "
  339. "`gpu_memory_utilization`, setting "
  340. "`--enable-chunked-prefill`, or `--kv-cache-dtype fp8` "
  341. "when initializing the engine. The last two are currently "
  342. "mutually exclusive.\n"
  343. f"Forcing max_model_len to {max_seq_len}.")