123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328 |
- import asyncio
- import os
- import pickle
- from collections import defaultdict
- from itertools import islice, repeat
- from typing import TYPE_CHECKING, Any, Dict, List, Optional, Tuple
- from aphrodite.common.sequence import ExecuteModelRequest, SamplerOutput
- from aphrodite.common.utils import (get_aphrodite_instance_id,
- get_distributed_init_method, get_ip,
- get_open_port, make_async)
- from aphrodite.executor.distributed_gpu_executor import (
- DistributedGPUExecutor, DistributedGPUExecutorAsync)
- from aphrodite.executor.ray_utils import RayWorkerWrapper, ray
- if ray is not None:
- from ray.util.scheduling_strategies import PlacementGroupSchedulingStrategy
- if TYPE_CHECKING:
- from ray.util.placement_group import PlacementGroup
- # If the env var is set, it uses the Ray's compiled DAG API
- # which optimizes the control plane overhead.
- # Run Aphrodite with APHRODITE_USE_RAY_COMPILED_DAG=1 to enable it.
- USE_RAY_COMPILED_DAG = bool(os.getenv("APHRODITE_USE_RAY_COMPILED_DAG", 0))
- class RayGPUExecutor(DistributedGPUExecutor):
- def _init_executor(self) -> None:
- assert self.parallel_config.distributed_executor_backend == "ray"
- placement_group = self.parallel_config.placement_group
- # Disable Ray usage stats collection.
- ray_usage = os.environ.get("RAY_USAGE_STATS_ENABLED", "0")
- if ray_usage != "1":
- os.environ["RAY_USAGE_STATS_ENABLED"] = "0"
- # Create the parallel GPU workers.
- self._init_workers_ray(placement_group)
- self.forward_dag = None
- if USE_RAY_COMPILED_DAG:
- self.forward_dag = self._compiled_ray_dag()
- self.extra_execute_model_run_workers_kwargs[
- "use_ray_compiled_dag"] = True
- def _configure_ray_workers_use_nsight(self,
- ray_remote_kwargs) -> Dict[str, Any]:
- # If nsight profiling is enabled, we need to set the profiling
- # configuration for the ray workers as runtime env.
- runtime_env = ray_remote_kwargs.setdefault("runtime_env", {})
- runtime_env.update({
- "nsight": {
- "t": "cuda,cudnn,cublas",
- "o": "'worker_process_%p'",
- "cuda-graph-trace": "node",
- }
- })
- return ray_remote_kwargs
- def _init_workers_ray(self, placement_group: "PlacementGroup",
- **ray_remote_kwargs):
- if self.parallel_config.tensor_parallel_size == 1:
- # For single GPU case, we use a ray worker with constrained memory.
- num_gpus = self.cache_config.gpu_memory_utilization
- else:
- # Otherwise, the ray workers are allocated with a full GPU.
- num_gpus = 1
- # The driver dummy worker does not actually use any resources.
- # It holds the resource for the driver worker.
- self.driver_dummy_worker: RayWorkerWrapper = None
- # The remaining workers are the actual ray actors.
- self.workers: List[RayWorkerWrapper] = []
- if self.parallel_config.ray_workers_use_nsight:
- ray_remote_kwargs = self._configure_ray_workers_use_nsight(
- ray_remote_kwargs)
- # Create the workers.
- driver_ip = get_ip()
- for bundle_id, bundle in enumerate(placement_group.bundle_specs):
- if not bundle.get("GPU", 0):
- continue
- scheduling_strategy = PlacementGroupSchedulingStrategy(
- placement_group=placement_group,
- placement_group_capture_child_tasks=True,
- placement_group_bundle_index=bundle_id,
- )
- if self.speculative_config is not None:
- worker_module_name = "aphrodite.spec_decode.spec_decode_worker"
- worker_class_name = "create_spec_worker"
- else:
- worker_module_name = "aphrodite.task_handler.worker"
- worker_class_name = "Worker"
- worker = ray.remote(
- num_cpus=0,
- num_gpus=num_gpus,
- scheduling_strategy=scheduling_strategy,
- **ray_remote_kwargs,
- )(RayWorkerWrapper).remote(
- worker_module_name=worker_module_name,
- worker_class_name=worker_class_name,
- trust_remote_code=self.model_config.trust_remote_code,
- )
- worker_ip = ray.get(worker.get_node_ip.remote())
- if worker_ip == driver_ip and self.driver_dummy_worker is None:
- # If the worker is on the same node as the driver, we use it
- # as the resource holder for the driver process.
- self.driver_dummy_worker = worker
- self.driver_worker = RayWorkerWrapper(
- worker_module_name=worker_module_name,
- worker_class_name=worker_class_name,
- trust_remote_code=self.model_config.trust_remote_code,
- )
- else:
- # Else, added to the list of workers.
- self.workers.append(worker)
- if self.driver_dummy_worker is None:
- raise ValueError(
- "Ray does not allocate any GPUs on the driver node. Consider "
- "adjusting the Ray placement group or running the driver on a "
- "GPU node.")
- # Get the set of GPU IDs used on each node.
- worker_node_and_gpu_ids = self._run_workers("get_node_and_gpu_ids",
- use_dummy_driver=True)
- node_workers = defaultdict(list)
- node_gpus = defaultdict(list)
- for i, (node_id, gpu_ids) in enumerate(worker_node_and_gpu_ids):
- node_workers[node_id].append(i)
- node_gpus[node_id].extend(gpu_ids)
- for node_id, gpu_ids in node_gpus.items():
- node_gpus[node_id] = sorted(gpu_ids)
- APHRODITE_INSTANCE_ID = get_aphrodite_instance_id()
- # Set environment variables for the driver and workers.
- all_args_to_update_environment_variables = [({
- "CUDA_VISIBLE_DEVICES":
- ",".join(map(str, node_gpus[node_id])),
- "APHRODITE_INSTANCE_ID":
- APHRODITE_INSTANCE_ID,
- "APHRODITE_TRACE_FUNCTION":
- os.getenv("APHRODITE_TRACE_FUNCTION", "0"),
- }, ) for (node_id, _) in worker_node_and_gpu_ids]
- self._run_workers("update_environment_variables",
- all_args=all_args_to_update_environment_variables)
- distributed_init_method = get_distributed_init_method(
- driver_ip, get_open_port())
- # Initialize the actual workers inside worker wrapper.
- init_worker_all_kwargs = [
- self._get_worker_kwargs(
- local_rank=node_workers[node_id].index(rank),
- rank=rank,
- distributed_init_method=distributed_init_method,
- ) for rank, (node_id, _) in enumerate(worker_node_and_gpu_ids)
- ]
- self._run_workers("init_worker", all_kwargs=init_worker_all_kwargs)
- self._run_workers("init_device")
- self._run_workers("load_model",
- max_concurrent_workers=self.parallel_config.
- max_parallel_loading_workers)
- def _driver_execute_model(
- self,
- execute_model_req: Optional[ExecuteModelRequest] = None
- ) -> List[SamplerOutput]:
- """Run execute_model in the driver worker.
- Passing None will cause the driver to stop the model execution
- loop running in each of the remote workers.
- """
- return self.driver_worker.execute_method("execute_model",
- execute_model_req)
- def _run_workers(
- self,
- method: str,
- *args,
- async_run_remote_workers_only: bool = False,
- all_args: Optional[List[Tuple[Any, ...]]] = None,
- all_kwargs: Optional[List[Dict[str, Any]]] = None,
- use_dummy_driver: bool = False,
- max_concurrent_workers: Optional[int] = None,
- use_ray_compiled_dag: bool = False,
- **kwargs,
- ) -> Any:
- """Runs the given method on all workers. Can be used in the following
- ways:
- - async_run_remote_workers_only: If True the method will be run only
- in the remote workers, not the driver worker. It will also be
- run asynchronously and return a list of futures rather than blocking
- on the results.
- - args/kwargs: All workers share the same args/kwargs
- - all_args/all_kwargs: args/kwargs for each worker are specified
- individually
- """
- if max_concurrent_workers:
- raise NotImplementedError(
- "max_concurrent_workers is not supported yet.")
- count = len(self.workers)
- all_worker_args = repeat(args, count) if all_args is None \
- else islice(all_args, 1, None)
- all_worker_kwargs = repeat(kwargs, count) if all_kwargs is None \
- else islice(all_kwargs, 1, None)
- if use_ray_compiled_dag:
- # Right now, compiled DAG can only accept a single
- # input. TODO: Fix it.
- assert self.forward_dag is not None
- output_channels = self.forward_dag.execute(1)
- ray_worker_outputs = []
- else:
- # Start the ray workers first.
- ray_worker_outputs = [
- worker.execute_method.remote(method, *worker_args,
- **worker_kwargs)
- for (worker, worker_args, worker_kwargs
- ) in zip(self.workers, all_worker_args, all_worker_kwargs)
- ]
- if async_run_remote_workers_only:
- # Just return futures
- return ray_worker_outputs
- driver_args = args if all_args is None else all_args[0]
- driver_kwargs = kwargs if all_kwargs is None else all_kwargs[0]
- # Start the driver worker after all the ray workers.
- if not use_dummy_driver:
- driver_worker_output = self.driver_worker.execute_method(
- method, *driver_args, **driver_kwargs)
- else:
- driver_worker_output = ray.get(
- self.driver_dummy_worker.execute_method.remote(
- method, *driver_args, **driver_kwargs))
- # Get the results of the ray workers.
- if self.workers:
- if use_ray_compiled_dag:
- try:
- ray_worker_outputs = [
- pickle.loads(chan.begin_read())
- for chan in output_channels
- ]
- finally:
- # Has to call end_read in order to reuse the DAG.
- for chan in output_channels:
- chan.end_read()
- else:
- ray_worker_outputs = ray.get(ray_worker_outputs)
- return [driver_worker_output] + ray_worker_outputs
- def _wait_for_tasks_completion(self, parallel_worker_tasks: Any) -> None:
- """Wait for futures returned from _run_workers() with
- async_run_remote_workers_only to complete."""
- ray.get(parallel_worker_tasks)
- def _compiled_ray_dag(self):
- import pkg_resources
- required_version = "2.9"
- current_version = pkg_resources.get_distribution("ray").version
- if current_version < required_version:
- raise ValueError(f"Ray version {required_version} or greater is "
- f"required, but found {current_version}")
- from ray.dag import InputNode, MultiOutputNode
- assert self.parallel_config.distributed_executor_backend == "ray"
- # Right now, compiled DAG requires at least 1 arg. We send
- # a dummy value for now. It will be fixed soon.
- with InputNode() as input_data:
- forward_dag = MultiOutputNode([
- worker.execute_model_compiled_dag_remote.bind(input_data)
- for worker in self.workers
- ])
- return forward_dag.experimental_compile()
- def check_health(self) -> None:
- """Raises an error if engine is unhealthy."""
- self._check_if_any_actor_is_dead()
- def _check_if_any_actor_is_dead(self):
- if not self.workers:
- return
- dead_actors = []
- for actor in self.workers:
- actor_state = ray.state.actors(actor._ray_actor_id.hex()) # pylint: disable=protected-access
- if actor_state["State"] == "DEAD":
- dead_actors.append(actor)
- if dead_actors:
- raise RuntimeError("At least one Worker is dead. "
- f"Dead Workers: {dead_actors}. ")
- class RayGPUExecutorAsync(RayGPUExecutor, DistributedGPUExecutorAsync):
- def __init__(self, *args, **kwargs):
- super().__init__(*args, **kwargs)
- self.driver_exec_method = make_async(self.driver_worker.execute_method)
- async def _driver_execute_model_async(
- self,
- execute_model_req: Optional[ExecuteModelRequest] = None
- ) -> List[SamplerOutput]:
- return await self.driver_exec_method("execute_model",
- execute_model_req)
- async def _start_worker_execution_loop(self):
- coros = [
- worker.execute_method.remote("start_worker_execution_loop")
- for worker in self.workers
- ]
- return await asyncio.gather(*coros)
|