llm_engine.py 12.6 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3

4
from collections.abc import Mapping
5
from copy import copy
6
from typing import Any, Callable, Optional, Union
7

8
9
from typing_extensions import TypeVar

10
import vllm.envs as envs
11
from vllm.config import ParallelConfig, VllmConfig
12
from vllm.distributed import stateless_destroy_torch_distributed_process_group
13
from vllm.engine.arg_utils import EngineArgs
14
from vllm.inputs import PromptType
15
16
from vllm.logger import init_logger
from vllm.lora.request import LoRARequest
17
from vllm.multimodal import MULTIMODAL_REGISTRY, MultiModalRegistry
18
from vllm.outputs import PoolingRequestOutput, RequestOutput
19
from vllm.pooling_params import PoolingParams
20
from vllm.sampling_params import SamplingParams
21
from vllm.tasks import SupportedTask
22
from vllm.transformers_utils.tokenizer_group import (
23
    TokenizerGroup, init_tokenizer_from_configs)
24
from vllm.usage.usage_lib import UsageContext
25
from vllm.utils import Device
26
from vllm.v1.engine.core_client import EngineCoreClient
27
from vllm.v1.engine.output_processor import OutputProcessor
28
from vllm.v1.engine.parallel_sampling import ParentRequest
29
from vllm.v1.engine.processor import Processor
30
from vllm.v1.executor.abstract import Executor
31
32
33
34
from vllm.v1.metrics.loggers import (PrometheusStatLogger, StatLoggerBase,
                                     StatLoggerFactory)
from vllm.v1.metrics.reader import Metric, get_metrics_snapshot
from vllm.v1.metrics.stats import IterationStats
35
36
37

logger = init_logger(__name__)

38
_R = TypeVar("_R", default=Any)
39

40
41

class LLMEngine:
42
    """Legacy LLMEngine for backwards compatibility."""
43
44
45

    def __init__(
        self,
46
        vllm_config: VllmConfig,
47
        executor_class: type[Executor],
48
49
        log_stats: bool,
        usage_context: UsageContext = UsageContext.ENGINE_CONTEXT,
50
        stat_loggers: Optional[list[StatLoggerFactory]] = None,
51
        mm_registry: MultiModalRegistry = MULTIMODAL_REGISTRY,
52
        use_cached_outputs: bool = False,
53
        multiprocess_mode: bool = False,
54
    ) -> None:
55
56
57
58
59
60
61
        if not envs.VLLM_USE_V1:
            raise ValueError(
                "Using V1 LLMEngine, but envs.VLLM_USE_V1=False. "
                "This should not happen. As a workaround, try using "
                "LLMEngine.from_vllm_config(...) or explicitly set "
                "VLLM_USE_V1=0 or 1 and report this issue on Github.")

62
63
64
65
66
        if stat_loggers is not None:
            raise NotImplementedError(
                "Passing StatLoggers to LLMEngine in V1 is not yet supported. "
                "Set VLLM_USE_V1=0 and file and issue on Github.")

67
        self.vllm_config = vllm_config
68
        self.model_config = vllm_config.model_config
69
        self.cache_config = vllm_config.cache_config
70

71
72
73
74
75
        self.log_stats = log_stats
        self.stat_logger: Optional[StatLoggerBase] = None
        if self.log_stats:
            self.stat_logger = PrometheusStatLogger(vllm_config)

76
        # important: init dp group before init the engine_core
77
78
79
80
81
82
        # In the decoupled engine case this is handled in EngineCoreProc.
        parallel_config = vllm_config.parallel_config
        if not multiprocess_mode and parallel_config.data_parallel_size > 1:
            self.dp_group = parallel_config.stateless_init_dp_group()
        else:
            self.dp_group = None
83
84
        self.should_execute_dummy_batch = False

85
86
87
88
89
90
91
92
        if self.model_config.skip_tokenizer_init:
            self.tokenizer = None
        else:
            # Tokenizer (+ ensure liveness if running in another process).
            self.tokenizer = init_tokenizer_from_configs(
                model_config=vllm_config.model_config,
                scheduler_config=vllm_config.scheduler_config,
                lora_config=vllm_config.lora_config)
93
94

        # Processor (convert Inputs --> EngineCoreRequests)
95
        self.processor = Processor(vllm_config=vllm_config,
96
97
                                   tokenizer=self.tokenizer,
                                   mm_registry=mm_registry)
98

99
100
        # OutputProcessor (convert EngineCoreOutputs --> RequestOutput).
        self.output_processor = OutputProcessor(self.tokenizer,
101
                                                log_stats=self.log_stats)
102
103
104
105
106

        # EngineCore (gets EngineCoreRequests and gives EngineCoreOutputs)
        self.engine_core = EngineCoreClient.make_client(
            multiprocess_mode=multiprocess_mode,
            asyncio_mode=False,
107
108
            vllm_config=vllm_config,
            executor_class=executor_class,
109
            log_stats=self.log_stats,
110
        )
111

112
113
114
115
        if not multiprocess_mode:
            # for v0 compatibility
            self.model_executor = self.engine_core.engine_core.model_executor  # type: ignore

116
117
118
        # Don't keep the dummy data in memory
        self.reset_mm_cache()

119
120
121
122
123
    @classmethod
    def from_vllm_config(
        cls,
        vllm_config: VllmConfig,
        usage_context: UsageContext = UsageContext.ENGINE_CONTEXT,
124
        stat_loggers: Optional[list[StatLoggerFactory]] = None,
125
126
127
128
129
130
131
132
133
        disable_log_stats: bool = False,
    ) -> "LLMEngine":
        return cls(vllm_config=vllm_config,
                   executor_class=Executor.get_class(vllm_config),
                   log_stats=(not disable_log_stats),
                   usage_context=usage_context,
                   stat_loggers=stat_loggers,
                   multiprocess_mode=envs.VLLM_ENABLE_V1_MULTIPROCESSING)

134
135
136
137
138
    @classmethod
    def from_engine_args(
        cls,
        engine_args: EngineArgs,
        usage_context: UsageContext = UsageContext.ENGINE_CONTEXT,
139
        stat_loggers: Optional[list[StatLoggerFactory]] = None,
140
        enable_multiprocessing: bool = False,
141
142
    ) -> "LLMEngine":
        """Creates an LLM engine from the engine arguments."""
143

144
        # Create the engine configs.
145
        vllm_config = engine_args.create_engine_config(usage_context)
146
        executor_class = Executor.get_class(vllm_config)
147

148
        if envs.VLLM_ENABLE_V1_MULTIPROCESSING:
149
150
151
152
153
154
155
156
157
158
159
160
            logger.debug("Enabling multiprocessing for LLMEngine.")
            enable_multiprocessing = True

        # Create the LLMEngine.
        return cls(vllm_config=vllm_config,
                   executor_class=executor_class,
                   log_stats=not engine_args.disable_log_stats,
                   usage_context=usage_context,
                   stat_loggers=stat_loggers,
                   multiprocess_mode=enable_multiprocessing)

    def get_num_unfinished_requests(self) -> int:
161
        return self.output_processor.get_num_unfinished_requests()
162
163

    def has_unfinished_requests(self) -> bool:
164
        has_unfinished = self.output_processor.has_unfinished_requests()
165
        if self.dp_group is None:
166
            return has_unfinished or self.engine_core.dp_engines_running()
167
168
169
170
171
172
173
174
        return self.has_unfinished_requests_dp(has_unfinished)

    def has_unfinished_requests_dp(self, has_unfinished: bool) -> bool:
        aggregated_has_unfinished = ParallelConfig.has_unfinished_dp(
            self.dp_group, has_unfinished)
        if not has_unfinished and aggregated_has_unfinished:
            self.should_execute_dummy_batch = True
        return aggregated_has_unfinished
175
176
177
178
179

    @classmethod
    def validate_outputs(cls, outputs, output_type):
        return outputs

180
181
182
    def get_supported_tasks(self) -> tuple[SupportedTask, ...]:
        return self.engine_core.get_supported_tasks()

183
    def abort_request(self, request_ids: list[str]) -> None:
184
185
        """Remove request_ids from EngineCore and Detokenizer."""

186
        request_ids = self.output_processor.abort_requests(request_ids)
187
188
        self.engine_core.abort_requests(request_ids)

189
190
191
192
193
194
195
    def add_request(
        self,
        request_id: str,
        prompt: PromptType,
        params: Union[SamplingParams, PoolingParams],
        arrival_time: Optional[float] = None,
        lora_request: Optional[LoRARequest] = None,
196
        tokenization_kwargs: Optional[dict[str, Any]] = None,
197
198
199
        trace_headers: Optional[Mapping[str, str]] = None,
        priority: int = 0,
    ) -> None:
200
201
202
203
204
        # Validate the request_id type.
        if not isinstance(request_id, str):
            raise TypeError(
                f"request_id must be a string, got {type(request_id)}")

205
        # Process raw inputs into the request.
206
207
        prompt_str, request = self.processor.process_inputs(
            request_id, prompt, params, arrival_time, lora_request,
208
            tokenization_kwargs, trace_headers, priority)
209

210
        n = params.n if isinstance(params, SamplingParams) else 1
211

212
213
        if n == 1:
            # Make a new RequestState and queue.
214
            self.output_processor.add_request(request, prompt_str, None, 0)
215
            # Add the request to EngineCore.
216
            self.engine_core.add_request(request)
217
218
219
220
221
222
223
224
225
226
227
            return

        # Fan out child requests (for n>1).
        parent_req = ParentRequest(request_id, params)
        for idx in range(n):
            request_id, params = parent_req.get_child_info(idx)
            child_request = request if idx == n - 1 else copy(request)
            child_request.request_id = request_id
            child_request.sampling_params = params

            # Make a new RequestState and queue.
228
229
            self.output_processor.add_request(child_request, prompt_str,
                                              parent_req, idx)
230
231
            # Add the request to EngineCore.
            self.engine_core.add_request(child_request)
232

233
    def step(self) -> Union[list[RequestOutput], list[PoolingRequestOutput]]:
234

235
236
237
238
239
        if self.should_execute_dummy_batch:
            self.should_execute_dummy_batch = False
            self.engine_core.execute_dummy_batch()
            return []

240
        # 1) Get EngineCoreOutput from the EngineCore.
241
        outputs = self.engine_core.get_output()
242

243
        # 2) Process EngineCoreOutputs.
244
        iteration_stats = IterationStats() if self.log_stats else None
245
        processed_outputs = self.output_processor.process_outputs(
246
247
248
            outputs.outputs,
            engine_core_timestamp=outputs.timestamp,
            iteration_stats=iteration_stats)
249

250
251
        # 3) Abort any reqs that finished due to stop strings.
        self.engine_core.abort_requests(processed_outputs.reqs_to_abort)
252

253
254
255
256
257
258
        # 4) Record stats
        if self.stat_logger is not None:
            assert outputs.scheduler_stats is not None
            self.stat_logger.record(scheduler_stats=outputs.scheduler_stats,
                                    iteration_stats=iteration_stats)

259
        return processed_outputs.request_outputs
260

261
262
263
    def get_vllm_config(self):
        return self.vllm_config

264
    def get_model_config(self):
265
        return self.model_config
266

267
    def start_profile(self):
268
        self.engine_core.profile(True)
269

270
    def stop_profile(self):
271
        self.engine_core.profile(False)
272

273
    def reset_mm_cache(self):
274
        self.processor.clear_cache()
275
276
        self.engine_core.reset_mm_cache()

277
    def reset_prefix_cache(self, device: Optional[Device] = None):
278
279
        self.engine_core.reset_prefix_cache()

280
281
282
    def sleep(self, level: int = 1):
        self.engine_core.sleep(level)

283
284
    def wake_up(self, tags: Optional[list[str]] = None):
        self.engine_core.wake_up(tags)
285

286
287
288
    def is_sleeping(self) -> bool:
        return self.engine_core.is_sleeping()

289
290
291
292
    def get_metrics(self) -> list[Metric]:
        assert self.log_stats, "Stat logging disabled"
        return get_metrics_snapshot()

293
294
    def get_tokenizer_group(self) -> TokenizerGroup:
        if self.tokenizer is None:
295
296
297
            raise ValueError("Unable to get tokenizer because "
                             "skip_tokenizer_init is True")

298
        return self.tokenizer
299
300
301
302
303
304
305
306
307

    def add_lora(self, lora_request: LoRARequest) -> bool:
        """Load a new LoRA adapter into the engine for future requests."""
        return self.engine_core.add_lora(lora_request)

    def remove_lora(self, lora_id: int) -> bool:
        """Remove an already loaded LoRA adapter."""
        return self.engine_core.remove_lora(lora_id)

308
    def list_loras(self) -> set[int]:
309
310
311
312
313
314
        """List all registered adapters."""
        return self.engine_core.list_loras()

    def pin_lora(self, lora_id: int) -> bool:
        """Prevent an adapter from being evicted."""
        return self.engine_core.pin_lora(lora_id)
315

316
317
318
319
320
321
322
    def collective_rpc(self,
                       method: Union[str, Callable[..., _R]],
                       timeout: Optional[float] = None,
                       args: tuple = (),
                       kwargs: Optional[dict[str, Any]] = None) -> list[_R]:
        return self.engine_core.collective_rpc(method, timeout, args, kwargs)

323
324
325
    def __del__(self):
        if dp_group := getattr(self, "dp_group", None):
            stateless_destroy_torch_distributed_process_group(dp_group)