"vscode:/vscode.git/clone" did not exist on "a388252ac49ec642bec5a8a63b226d3058e64a54"
protocol.py 11.7 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3

4
5
import asyncio
from abc import ABC, abstractmethod
6
from typing import Any, AsyncGenerator, Iterable, Mapping, Optional, Union
7

8
from vllm.beam_search import BeamSearchSequence, create_sort_beams_key_function
9
from vllm.config import ModelConfig, VllmConfig
10
from vllm.inputs.data import PromptType, TokensPrompt
11
from vllm.inputs.parse import is_explicit_encoder_decoder_prompt
12
from vllm.inputs.preprocess import InputPreprocessor
13
from vllm.logger import init_logger
14
from vllm.lora.request import LoRARequest
15
from vllm.outputs import CompletionOutput, PoolingRequestOutput, RequestOutput
16
from vllm.plugins.io_processors.interface import IOProcessor
17
from vllm.pooling_params import PoolingParams
18
from vllm.sampling_params import BeamSearchParams, SamplingParams
19
from vllm.tasks import SupportedTask
20
from vllm.transformers_utils.tokenizer import AnyTokenizer
21
from vllm.utils import Device, collect_from_async_generator, random_uuid
22

23
logger = init_logger(__name__)
24

25
26

class EngineClient(ABC):
27
    """Protocol class for Clients to Engine"""
28
29

    @property
30
    @abstractmethod
31
32
33
34
    def is_running(self) -> bool:
        ...

    @property
35
    @abstractmethod
36
37
38
39
    def is_stopped(self) -> bool:
        ...

    @property
40
    @abstractmethod
41
42
43
    def errored(self) -> bool:
        ...

44
    @property
45
    @abstractmethod
46
47
    def dead_error(self) -> BaseException:
        ...
48

49
    @abstractmethod
50
    def generate(
51
        self,
52
        prompt: PromptType,
53
54
55
56
        sampling_params: SamplingParams,
        request_id: str,
        lora_request: Optional[LoRARequest] = None,
        trace_headers: Optional[Mapping[str, str]] = None,
57
        priority: int = 0,
58
    ) -> AsyncGenerator[RequestOutput, None]:
59
        """Generate outputs for a request."""
60
        ...
61

62
63
    async def beam_search(
        self,
64
        prompt: PromptType,
65
66
        request_id: str,
        params: BeamSearchParams,
67
        lora_request: Optional[LoRARequest] = None,
68
69
70
71
72
73
74
    ) -> AsyncGenerator[RequestOutput, None]:

        beam_width = params.beam_width
        max_tokens = params.max_tokens
        ignore_eos = params.ignore_eos
        temperature = params.temperature
        length_penalty = params.length_penalty
75
        include_stop_str_in_output = params.include_stop_str_in_output
76

77
        preprocessor = await self.get_input_preprocessor()
78
        tokenizer = preprocessor.get_tokenizer()
79
        eos_token_id = tokenizer.eos_token_id
80

81
82
83
        if is_explicit_encoder_decoder_prompt(prompt):
            raise NotImplementedError
        else:
84
            processed_inputs = preprocessor._prompt_to_llm_inputs(prompt)
85

86
87
88
        if processed_inputs["type"] == "embeds":
            raise NotImplementedError

89
90
91
92
93
94
95
96
97
98
99
        # This is a workaround to fix multimodal beam search; this is a
        # bandaid fix for 2 small problems:
        # 1. Multi_modal_data on the processed_inputs currently resolves to
        #    `None`.
        # 2. preprocessing above expands the multimodal placeholders. However,
        #    this happens again in generation, so the double expansion causes
        #    a mismatch.
        # TODO - would be ideal to handle this more gracefully.
        prompt_token_ids = prompt.get("prompt_token_ids")
        multi_modal_data = prompt.get("multi_modal_data")

100
101
102
        prompt_text = processed_inputs.get("prompt")
        mm_processor_kwargs = processed_inputs.get("mm_processor_kwargs")

103
        tokenized_length = len(prompt_token_ids)
104
105

        sort_beams_key = create_sort_beams_key_function(
106
            eos_token_id, length_penalty)
107

108
109
110
111
112
        beam_search_params = SamplingParams(
            logprobs=2 * beam_width,
            max_tokens=1,
            temperature=temperature,
        )
113
        all_beams = [
114
115
            BeamSearchSequence(tokens=prompt_token_ids,
                               cum_logprob=0,
116
                               logprobs=[],
117
                               multi_modal_data=multi_modal_data,
118
119
                               mm_processor_kwargs=mm_processor_kwargs,
                               lora_request=lora_request)
120
        ]
121
122
123
        completed = []

        for _ in range(max_tokens):
124
            prompts_batch, lora_req_batch = zip(*[(
125
126
                TokensPrompt(prompt_token_ids=beam.tokens,
                             multi_modal_data=beam.multi_modal_data,
127
128
129
                             mm_processor_kwargs=beam.mm_processor_kwargs),
                beam.lora_request,
            ) for beam in all_beams])
130
131
132
133

            tasks = []

            request_id = f"beam_search-{random_uuid()}"
134
135
            for i, (individual_prompt,
                    lora_req) in enumerate(zip(prompts_batch, lora_req_batch)):
136
137
138
                request_id_item = f"{request_id}-{i}"
                task = asyncio.create_task(
                    collect_from_async_generator(
139
140
141
142
                        self.generate(individual_prompt,
                                      beam_search_params,
                                      request_id_item,
                                      lora_request=lora_req)))
143
144
145
146
147
148
149
150
151
152
153
154
155
                tasks.append(task)

            output = await asyncio.gather(*tasks)

            output = [x[0] for x in output]

            new_beams = []
            for i, current_beam in enumerate(all_beams):
                result = output[i]

                if result.outputs[0].logprobs is not None:
                    logprobs = result.outputs[0].logprobs[0]
                    for token_id, logprob_obj in logprobs.items():
156
                        if token_id == eos_token_id and \
157
                            not ignore_eos:
158
159
160
161
162
163
164
165
166
167
                            completed.append(
                                BeamSearchSequence(
                                    tokens=current_beam.tokens +
                                    [token_id] if include_stop_str_in_output
                                    else current_beam.tokens,
                                    logprobs=current_beam.logprobs +
                                    [logprobs],
                                    cum_logprob=current_beam.cum_logprob +
                                    logprob_obj.logprob,
                                    finish_reason="stop",
168
                                    stop_reason=eos_token_id))
169
                        else:
170
171
172
173
174
                            new_beams.append(
                                BeamSearchSequence(
                                    tokens=current_beam.tokens + [token_id],
                                    logprobs=current_beam.logprobs +
                                    [logprobs],
175
                                    lora_request=current_beam.lora_request,
176
177
178
179
180
181
                                    cum_logprob=current_beam.cum_logprob +
                                    logprob_obj.logprob,
                                    multi_modal_data=current_beam.
                                    multi_modal_data,
                                    mm_processor_kwargs=current_beam.
                                    mm_processor_kwargs))
182
183
184
185
186
187
188
189
190

            sorted_beams = sorted(new_beams, key=sort_beams_key, reverse=True)
            all_beams = sorted_beams[:beam_width]

        completed.extend(all_beams)
        sorted_completed = sorted(completed, key=sort_beams_key, reverse=True)
        best_beams = sorted_completed[:beam_width]

        for beam in best_beams:
191
            if (beam.tokens[-1] == eos_token_id and not ignore_eos):
Robert Shaw's avatar
Robert Shaw committed
192
193
194
195
196
                # Skip the eos token in the text.
                tokens = beam.tokens[tokenized_length:-1]
            else:
                tokens = beam.tokens[tokenized_length:]
            beam.text = tokenizer.decode(tokens)
197

198
        yield RequestOutput(
199
            request_id=request_id,
200
            prompt=prompt_text,
201
            outputs=[
202
203
204
205
206
207
208
209
210
                CompletionOutput(text=beam.text,
                                 cumulative_logprob=beam.cum_logprob,
                                 token_ids=beam.tokens[tokenized_length:],
                                 index=i,
                                 logprobs=beam.logprobs,
                                 finish_reason=beam.finish_reason if
                                 beam.finish_reason is not None else "length",
                                 stop_reason=beam.stop_reason)
                for (i, beam) in enumerate(best_beams)
211
212
            ],
            finished=True,
213
            prompt_token_ids=prompt_token_ids,
214
215
216
            prompt_logprobs=None)

    @abstractmethod
217
    def encode(
218
        self,
219
        prompt: PromptType,
220
221
222
223
        pooling_params: PoolingParams,
        request_id: str,
        lora_request: Optional[LoRARequest] = None,
        trace_headers: Optional[Mapping[str, str]] = None,
224
        priority: int = 0,
225
        tokenization_kwargs: Optional[dict[str, Any]] = None,
226
    ) -> AsyncGenerator[PoolingRequestOutput, None]:
227
        """Generate outputs for a request from a pooling model."""
228
        ...
229

230
    @abstractmethod
231
    async def abort(self, request_id: Union[str, Iterable[str]]) -> None:
232
233
234
        """Abort a request.

        Args:
235
236
            request_id: The unique id of the request,
                        or an iterable of such ids.
237
        """
238
        ...
239
240
241
242
243

    @abstractmethod
    async def get_vllm_config(self) -> VllmConfig:
        """Get the vllm configuration of the vLLM engine."""
        ...
244

245
    @abstractmethod
246
247
    async def get_model_config(self) -> ModelConfig:
        """Get the model configuration of the vLLM engine."""
248
        ...
249

250
251
252
253
    @abstractmethod
    async def get_input_preprocessor(self) -> InputPreprocessor:
        """Get the input processor of the vLLM engine."""
        ...
254

255
    @abstractmethod
256
257
    async def get_tokenizer(self) -> AnyTokenizer:
        """Get the tokenizer"""
258
        ...
259

260
261
262
    async def get_io_processor(self) -> IOProcessor:
        raise NotImplementedError

263
    @abstractmethod
264
    async def is_tracing_enabled(self) -> bool:
265
        ...
266

267
    @abstractmethod
268
    async def do_log_stats(self) -> None:
269
        ...
270

271
    @abstractmethod
272
273
    async def check_health(self) -> None:
        """Raise if unhealthy"""
274
        ...
275

276
    @abstractmethod
277
278
279
280
    async def start_profile(self) -> None:
        """Start profiling the engine"""
        ...

281
    @abstractmethod
282
283
284
    async def stop_profile(self) -> None:
        """Start profiling the engine"""
        ...
285

286
287
288
289
290
    @abstractmethod
    async def reset_mm_cache(self) -> None:
        """Reset the multi-modal cache"""
        ...

291
    @abstractmethod
292
293
    async def reset_prefix_cache(self,
                                 device: Optional[Device] = None) -> None:
294
295
296
        """Reset the prefix cache"""
        ...

297
298
299
300
301
302
    @abstractmethod
    async def sleep(self, level: int = 1) -> None:
        """Sleep the engine"""
        ...

    @abstractmethod
303
    async def wake_up(self, tags: Optional[list[str]] = None) -> None:
304
305
306
        """Wake up the engine"""
        ...

307
308
309
310
311
    @abstractmethod
    async def is_sleeping(self) -> bool:
        """Check whether the engine is sleeping"""
        ...

312
    @abstractmethod
313
    async def add_lora(self, lora_request: LoRARequest) -> bool:
314
315
        """Load a new LoRA adapter into the engine for future requests."""
        ...
316
317
318
319
320
321

    async def scale_elastic_ep(self,
                               new_data_parallel_size: int,
                               drain_timeout: int = 300) -> None:
        """Scale the engine"""
        raise NotImplementedError
322
323
324
325
326
327
328
329

    async def collective_rpc(self,
                             method: str,
                             timeout: Optional[float] = None,
                             args: tuple = (),
                             kwargs: Optional[dict] = None):
        """Perform a collective RPC call to the given path."""
        raise NotImplementedError
330
331
332
333

    async def get_supported_tasks(self) -> tuple[SupportedTask, ...]:
        """Get supported tasks"""
        raise NotImplementedError