preprocess.py 18.2 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 typing import Any, overload
6
7
8

from typing_extensions import assert_never

9
from vllm.config import VllmConfig
10
from vllm.logger import init_logger
11
from vllm.multimodal import MULTIMODAL_REGISTRY, MultiModalRegistry
12
from vllm.multimodal.cache import BaseMultiModalProcessorCache
13
14
15
16
17
from vllm.multimodal.inputs import (
    MultiModalDataDict,
    MultiModalInputs,
    MultiModalUUIDDict,
)
18
from vllm.multimodal.processing import BaseMultiModalProcessor
19
from vllm.renderers import BaseRenderer, renderer_from_config
20
21
22
23
24
25
26
27
28
29
from vllm.renderers.inputs import (
    DecoderDictPrompt,
    DecoderOnlyDictPrompt,
    DictPrompt,
    EncoderDecoderDictPrompt,
    EncoderDictPrompt,
    SingletonDictPrompt,
    TokPrompt,
)
from vllm.renderers.inputs.preprocess import parse_dec_only_prompt, parse_enc_dec_prompt
30
from vllm.tokenizers import TokenizerLike
31
from vllm.utils.jsontree import json_iter_leaves
32
from vllm.v1.metrics.stats import MultiModalCacheStats
33

34
from .data import (
35
    DecoderInputs,
36
37
38
39
    DecoderOnlyInputs,
    EmbedsInputs,
    EmbedsPrompt,
    EncoderDecoderInputs,
40
    EncoderInputs,
41
42
43
44
45
46
47
48
49
    ProcessorInputs,
    PromptType,
    SingletonInputs,
    TextPrompt,
    TokenInputs,
    TokensPrompt,
    embeds_inputs,
    token_inputs,
)
50
51
52
53
54
55
56

logger = init_logger(__name__)


class InputPreprocessor:
    def __init__(
        self,
57
        vllm_config: VllmConfig,
58
        renderer: BaseRenderer | None = None,
59
        mm_registry: MultiModalRegistry = MULTIMODAL_REGISTRY,
60
        mm_processor_cache: BaseMultiModalProcessorCache | None = None,
61
62
63
    ) -> None:
        super().__init__()

64
65
66
        self.model_config = vllm_config.model_config
        self.observability_config = vllm_config.observability_config
        self.renderer = renderer or renderer_from_config(vllm_config)
67
        self.mm_registry = mm_registry
68
        self.mm_processor_cache = mm_processor_cache
69

70
71
        self.mm_cache_stats = MultiModalCacheStats() if mm_processor_cache else None

72
73
74
    @property
    def tokenizer(self) -> TokenizerLike | None:
        return self.renderer.tokenizer
75

76
77
    def get_tokenizer(self) -> TokenizerLike:
        return self.renderer.get_tokenizer()
78

79
    def get_decoder_start_token_id(self) -> int:
80
        """
81
        Obtain the decoder start token id employed by an encoder/decoder
82
        model. Raises an error if it is not available.
83
        """
84
85
86
        dec_start_token_id = getattr(
            self.model_config.hf_config, "decoder_start_token_id", None
        )
87

88
        if dec_start_token_id is None:
89
            logger.warning_once(
90
91
                "Falling back on <BOS> for decoder start token id "
                "because decoder start token id is not available."
92
            )
93
            dec_start_token_id = self.renderer.get_bos_token_id()
94

95
96
        if dec_start_token_id is None:
            raise RuntimeError("Cannot find decoder start token id or <BOS>")
97

98
        return dec_start_token_id
99

100
    def _prepare_decoder_input_ids(self, decoder_input_ids: list[int]) -> list[int]:
101
102
103
        """
        Prepares `decoder_input_ids` for generation with encoder-decoder models.

104
105
106
107
        Based on:
        https://github.com/huggingface/transformers/blob/4037a2b5b1278736e566aec12e169100275545ea/src/transformers/generation/utils.py
        specifically,
        `GenerationMixin._prepare_decoder_input_ids_for_generation()`.
108
109
110
111
112
113
114
115
116
117
118

        Arguments:

        * decoder_input_ids: input token ids to preprocess

        Returns:

        * Processed token list
        """
        decoder_start_token_id = self.get_decoder_start_token_id()

119
120
121
122
        if (
            len(decoder_input_ids) == 0
            or decoder_input_ids[0] != decoder_start_token_id
        ):
123
124
125
126
            decoder_input_ids = [decoder_start_token_id] + decoder_input_ids

        return decoder_input_ids

127
128
    def _get_tokenization_kw(
        self,
129
        overrides: dict[str, Any] | None = None,
130
131
132
    ) -> dict[str, Any]:
        kwargs = dict[str, Any]()

133
        if self.model_config.is_encoder_decoder:
134
135
136
137
138
139
140
141
142
143
            # For Whisper, special tokens should be provided by the user based
            # on the task and language of their request. Also needed to avoid
            # appending an EOS token to the prompt which disrupts generation.
            kwargs["add_special_tokens"] = False

        if overrides:
            kwargs.update(overrides)

        return kwargs

144
145
146
    def _tokenize_prompt(
        self,
        prompt: str,
147
        tokenization_kwargs: dict[str, Any] | None = None,
148
    ) -> list[int]:
149
150
151
152
        """
        Apply the model's tokenizer to a text prompt, returning the
        corresponding token IDs.
        """
153
        tokenizer = self.get_tokenizer()
154
        tokenization_kwargs = self._get_tokenization_kw(tokenization_kwargs)
155

156
        encoder_config = self.model_config.encoder_config
157

158
        if encoder_config and encoder_config.get("do_lower_case", False):
159
160
            prompt = prompt.lower()

161
        return tokenizer.encode(prompt, **tokenization_kwargs)
162

163
164
165
    def _get_mm_processor(self) -> BaseMultiModalProcessor:
        if not hasattr(self, "_mm_processor"):
            self._mm_processor = self.mm_registry.create_processor(
166
                self.model_config,
167
                self.observability_config,
168
                tokenizer=self.tokenizer,
169
170
171
172
                cache=self.mm_processor_cache,
            )

        return self._mm_processor
173

174
175
    def _process_multimodal(
        self,
176
        prompt: str | list[int],
177
        mm_data: MultiModalDataDict,
178
179
        mm_processor_kwargs: Mapping[str, object] | None,
        tokenization_kwargs: dict[str, Any] | None = None,
180
        *,
181
        mm_uuids: MultiModalUUIDDict | None = None,
182
    ) -> MultiModalInputs:
183
184
185
186
        """
        Apply the model's multi-modal processor to a multi-modal prompt,
        returning the corresponding token IDs and metadata.
        """
187
        mm_processor = self._get_mm_processor()
188

189
190
191
        if mm_processor_kwargs is None:
            mm_processor_kwargs = {}

192
        mm_items = mm_processor.info.parse_mm_data(mm_data)
193
        mm_input = mm_processor.apply(
194
            prompt,
195
            mm_items,
196
197
            hf_processor_mm_kwargs=mm_processor_kwargs,
            tokenization_kwargs=tokenization_kwargs,
198
            mm_uuids=mm_uuids,
199
        )
200
201
202
        mm_hashes = mm_input["mm_hashes"]

        # Validate that all mm items have a string as their hash
203
204
205
206
        contains_only_strings = all(
            isinstance(leaf, str) for leaf in json_iter_leaves(mm_hashes)
        )
        if not contains_only_strings:
207
208
209
            raise ValueError(
                f"mm_hashes must contain only strings, got: {mm_hashes}. "
                "This is likely due to an incorrect custom implementation of "
210
211
                "MultiModalProcessor.apply method."
            )
212
213

        return mm_input
214

215
216
217
218
    def _process_embeds(
        self,
        parsed_content: EmbedsPrompt,
    ) -> EmbedsInputs:
219
        if not self.model_config.enable_prompt_embeds:
220
221
222
            raise ValueError(
                "You must set `--enable-prompt-embeds` to input `prompt_embeds`."
            )
223
224

        prompt_embeds = parsed_content["prompt_embeds"]
225

226
227
228
229
230
231
232
233
        # prompt_embeds must be (seq_len, hidden_size), but if the user
        # passes in a batch of size 1, i.e. (1, seq_len, hidden_size),
        # we can unambiguously process the intent by squeezing the batch
        # dimension.
        if prompt_embeds.ndim == 3:
            prompt_embeds = prompt_embeds.squeeze(dim=0)

        if prompt_embeds.ndim != 2:
234
            raise ValueError("prompt_embeds must be of shape (seq_len, hidden_size).")
235

236
237
238
239
240
        # Tensors must be on CPU for serialization between processes
        # in the MsgpackEncoder. Casting to CPU here ensures that there is no
        # hidden device transfer in the critical path of generation.
        prompt_embeds = prompt_embeds.cpu()

241
242
243
        return embeds_inputs(
            prompt_embeds=prompt_embeds, cache_salt=parsed_content.get("cache_salt")
        )
244

245
    def _truncate_inputs(
246
        self, inputs: list[int], tokenization_kwargs: dict[str, Any] | None = None
247
248
249
250
251
252
    ) -> list[int]:
        if (
            not tokenization_kwargs
            or "truncation" not in tokenization_kwargs
            or self.tokenizer is None
        ):
253
254
255
256
257
258
259
260
261
            return inputs

        max_length = tokenization_kwargs["max_length"]

        if self.tokenizer.truncation_side == "left":
            return inputs[-max_length:]
        else:
            return inputs[:max_length]

262
263
264
    def _process_tokens(
        self,
        parsed_content: TokensPrompt,
265
        tokenization_kwargs: dict[str, Any] | None = None,
266
        *,
267
268
        mm_uuids: MultiModalUUIDDict | None = None,
    ) -> TokenInputs | MultiModalInputs:
269
        prompt_token_ids = self._truncate_inputs(
270
271
            parsed_content["prompt_token_ids"], tokenization_kwargs
        )
272

273
        inputs: TokenInputs | MultiModalInputs
274
        if multi_modal_data := parsed_content.get("multi_modal_data"):
275
276
            inputs = self._process_multimodal(
                prompt_token_ids,
277
                multi_modal_data,
278
                parsed_content.get("mm_processor_kwargs") or {},
279
                tokenization_kwargs=tokenization_kwargs,
280
                mm_uuids=mm_uuids,
281
            )
282
        else:
283
            inputs = token_inputs(prompt_token_ids)
284
285
286
287
288
289
290
291
292

        if cache_salt := parsed_content.get("cache_salt"):
            inputs["cache_salt"] = cache_salt

        return inputs

    def _process_text(
        self,
        parsed_content: TextPrompt,
293
        tokenization_kwargs: dict[str, Any] | None = None,
294
        *,
295
296
        mm_uuids: MultiModalUUIDDict | None = None,
    ) -> TokenInputs | MultiModalInputs:
297
298
        prompt_text = parsed_content["prompt"]

299
        inputs: TokenInputs | MultiModalInputs
300
        if multi_modal_data := parsed_content.get("multi_modal_data"):
301
302
            inputs = self._process_multimodal(
                prompt_text,
303
                multi_modal_data,
304
                parsed_content.get("mm_processor_kwargs") or {},
305
                tokenization_kwargs=tokenization_kwargs,
306
                mm_uuids=mm_uuids,
307
308
309
310
311
312
            )
        else:
            prompt_token_ids = self._tokenize_prompt(
                prompt_text,
                tokenization_kwargs=tokenization_kwargs,
            )
313
            inputs = token_inputs(prompt_token_ids)
314
315
316
317
318

        if cache_salt := parsed_content.get("cache_salt"):
            inputs["cache_salt"] = cache_salt

        return inputs
319

320
    @overload
321
    def _prompt_to_llm_inputs(
322
        self,
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
        prompt: EncoderDictPrompt,
        tokenization_kwargs: dict[str, Any] | None = None,
        *,
        mm_uuids: MultiModalUUIDDict | None = None,
    ) -> EncoderInputs: ...

    @overload
    def _prompt_to_llm_inputs(  # type: ignore[misc]
        self,
        prompt: DecoderDictPrompt,
        tokenization_kwargs: dict[str, Any] | None = None,
        *,
        mm_uuids: MultiModalUUIDDict | None = None,
    ) -> DecoderInputs: ...

    @overload
    def _prompt_to_llm_inputs(  # type: ignore[misc]
        self,
        prompt: DecoderOnlyDictPrompt,
        tokenization_kwargs: dict[str, Any] | None = None,
        *,
        mm_uuids: MultiModalUUIDDict | None = None,
    ) -> DecoderOnlyInputs: ...

    def _prompt_to_llm_inputs(
        self,
        prompt: SingletonDictPrompt,
350
        tokenization_kwargs: dict[str, Any] | None = None,
351
        *,
352
        mm_uuids: MultiModalUUIDDict | None = None,
353
    ) -> SingletonInputs:
354
355
        """
        Extract the singleton inputs from a prompt.
356
357
358

        Arguments:

359
        * prompt: single encoder or decoder input prompt
360
361
362

        Returns:

363
        * [`SingletonInputs`][vllm.inputs.data.SingletonInputs] instance
364
        """
365
366
        if "prompt_embeds" in prompt:
            return self._process_embeds(prompt)  # type: ignore[arg-type]
367

368
        if "prompt_token_ids" in prompt:
369
            return self._process_tokens(
370
                prompt,  # type: ignore[arg-type]
371
                mm_uuids=mm_uuids,
372
            )
373
374

        if "prompt" in prompt:
375
            return self._process_text(
376
                prompt,  # type: ignore[arg-type]
377
                tokenization_kwargs=tokenization_kwargs,
378
                mm_uuids=mm_uuids,
379
            )
380

381
        assert_never(prompt)  # type: ignore[arg-type]
382

383
    def _validate_enc_inputs(self, inputs: SingletonInputs) -> EncoderInputs:
384
        if inputs["type"] == "embeds":
385
386
387
            raise ValueError(
                "Embedding inputs are not supported for encoder-decoder models"
            )
388

389
390
391
392
        if inputs["type"] == "multimodal" and "encoder_prompt_token_ids" not in inputs:
            raise RuntimeError(
                "You should register an encoder-decoder "
                "multi-modal processor for encoder-decoder models."
393
            )
394

395
        return inputs  # type: ignore[return-value]
396

397
    def _validate_dec_inputs(self, inputs: SingletonInputs) -> DecoderInputs:
398
        if inputs["type"] == "embeds":
399
400
401
            raise ValueError(
                "Embedding inputs are not supported for encoder-decoder models"
            )
402

403
        return inputs
404

405
406
407
408
409
410
    def _build_enc_dec_inputs(
        self,
        encoder_inputs: SingletonInputs,
        decoder_inputs: SingletonInputs | None = None,
    ) -> EncoderDecoderInputs:
        enc_inputs = self._validate_enc_inputs(encoder_inputs)
411

412
413
414
415
416
417
418
        if decoder_inputs is None:
            dec_inputs: DecoderInputs = enc_inputs  # type: ignore[assignment]
        else:
            dec_inputs = self._validate_dec_inputs(decoder_inputs)

        enc_inputs_new: EncoderInputs
        dec_inputs_new: DecoderInputs
419

420
421
422
        if enc_inputs["type"] == "multimodal":
            enc_inputs_new = token_inputs(enc_inputs["encoder_prompt_token_ids"])
            dec_inputs_new = MultiModalInputs(
423
                type="multimodal",
424
425
426
427
                prompt_token_ids=dec_inputs["prompt_token_ids"],
                mm_kwargs=enc_inputs["mm_kwargs"],
                mm_hashes=enc_inputs["mm_hashes"],
                mm_placeholders=enc_inputs["mm_placeholders"],
428
            )
429
430
431
        elif enc_inputs["type"] == "token":
            enc_inputs_new = token_inputs(prompt_token_ids=[])
            dec_inputs_new = dec_inputs
432
        else:
433
434
435
436
437
438
439
            assert_never(enc_inputs)

        dec_inputs_new["prompt_token_ids"] = self._prepare_decoder_input_ids(
            dec_inputs_new["prompt_token_ids"]
        )
        if cache_salt := enc_inputs.get("cache_salt"):
            dec_inputs_new["cache_salt"] = cache_salt
440

441
        return EncoderDecoderInputs(encoder=enc_inputs_new, decoder=dec_inputs_new)
442

443
444
    def _process_encoder_decoder_prompt(
        self,
445
        prompt: EncoderDecoderDictPrompt,
446
        tokenization_kwargs: dict[str, Any] | None = None,
447
        *,
448
        mm_uuids: MultiModalUUIDDict | None = None,
449
    ) -> EncoderDecoderInputs:
450
        """
451
        For encoder/decoder models only:
452
453
454
        Process an input prompt into an
        [`EncoderDecoderInputs`][vllm.inputs.data.EncoderDecoderInputs]
        instance.
455
456
457

        Arguments:

458
        * prompt: an input prompt
459
460
461

        Returns:

462
463
        * [`EncoderDecoderInputs`][vllm.inputs.data.EncoderDecoderInputs]
          instance
464
        """
465
466
        encoder_prompt = prompt["encoder_prompt"]
        decoder_prompt = prompt["decoder_prompt"]
467
468
469
470

        return self._build_enc_dec_inputs(
            encoder_inputs=self._prompt_to_llm_inputs(
                encoder_prompt,
471
                tokenization_kwargs=tokenization_kwargs,
472
                mm_uuids=mm_uuids,
473
474
475
476
477
478
479
            ),
            decoder_inputs=(
                None
                if decoder_prompt is None
                else self._prompt_to_llm_inputs(
                    decoder_prompt,
                    tokenization_kwargs=tokenization_kwargs,
480
                )
481
482
            ),
        )
483
484
485

    def _process_decoder_only_prompt(
        self,
486
        prompt: DecoderOnlyDictPrompt,
487
        tokenization_kwargs: dict[str, Any] | None = None,
488
        *,
489
        mm_uuids: MultiModalUUIDDict | None = None,
490
    ) -> DecoderOnlyInputs:
491
        """
492
        For decoder-only models:
493
494
        Process an input prompt into a
        [`DecoderOnlyInputs`][vllm.inputs.data.DecoderOnlyInputs] instance.
495
496
497

        Arguments:

498
        * prompt: input prompt
499
500
501

        Returns:

502
        * [`DecoderOnlyInputs`][vllm.inputs.data.DecoderOnlyInputs] instance
503
        """
504
        return self._prompt_to_llm_inputs(
505
            prompt,
506
            tokenization_kwargs=tokenization_kwargs,
507
            mm_uuids=mm_uuids,
508
509
        )

510
    def _preprocess(
511
        self,
512
        prompt: PromptType | DictPrompt | TokPrompt,
513
        tokenization_kwargs: dict[str, Any] | None = None,
514
        *,
515
        mm_uuids: MultiModalUUIDDict | None = None,
516
    ) -> ProcessorInputs:
517
        if self.model_config.is_encoder_decoder:
518
            # Encoder-decoder model requires special mapping of
519
            # input prompts to encoder & decoder.
520
            return self._process_encoder_decoder_prompt(
521
                parse_enc_dec_prompt(prompt),
522
                tokenization_kwargs,
523
                mm_uuids=mm_uuids,
524
            )
525
526

        return self._process_decoder_only_prompt(
527
            parse_dec_only_prompt(prompt),
528
            tokenization_kwargs=tokenization_kwargs,
529
            mm_uuids=mm_uuids,
530
531
        )

532
533
    def preprocess(
        self,
534
        prompt: PromptType | DictPrompt | TokPrompt,
535
        tokenization_kwargs: dict[str, Any] | None = None,
536
        *,
537
        mm_uuids: MultiModalUUIDDict | None = None,
538
539
    ) -> ProcessorInputs:
        """Preprocess the input prompt."""
540
        res = self._preprocess(prompt, tokenization_kwargs, mm_uuids=mm_uuids)
541
542
543
544
545
546
547
548
549

        if self.mm_processor_cache and self.mm_cache_stats is not None:
            delta = self.mm_processor_cache.make_stats(delta=True)
            self.mm_cache_stats.requests += 1
            self.mm_cache_stats.queries += delta.total
            self.mm_cache_stats.hits += delta.hits

        return res

550
    def stat_mm_cache(self) -> MultiModalCacheStats | None:
551
552
553
554
555
556
557
558
559
        mm_cache_stats = self.mm_cache_stats
        if mm_cache_stats is None:
            return None

        self.mm_cache_stats = MultiModalCacheStats()

        return mm_cache_stats

    def clear_mm_cache(self) -> None:
560
561
        if self.mm_processor_cache is not None:
            self.mm_processor_cache.clear_cache()
562
563
564

        if self.mm_cache_stats is not None:
            self.mm_cache_stats.reset = True