preprocess.py 18.3 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 ModelConfig, ObservabilityConfig
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
        model_config: ModelConfig,
58
        observability_config: ObservabilityConfig | None = None,
59
        renderer: BaseRenderer | None = None,
60
        mm_registry: MultiModalRegistry = MULTIMODAL_REGISTRY,
61
        mm_processor_cache: BaseMultiModalProcessorCache | None = None,
62
63
64
    ) -> None:
        super().__init__()

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

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

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

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

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

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

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

99
        return dec_start_token_id
100

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

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

        Arguments:

        * decoder_input_ids: input token ids to preprocess

        Returns:

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

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

        return decoder_input_ids

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

134
        if self.model_config.is_encoder_decoder:
135
136
137
138
139
140
141
142
143
144
            # 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

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

157
        encoder_config = self.model_config.encoder_config
158

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

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

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

        return self._mm_processor
174

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

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

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

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

        return mm_input
215

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

        prompt_embeds = parsed_content["prompt_embeds"]
226

227
228
229
230
231
232
233
234
        # 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:
235
            raise ValueError("prompt_embeds must be of shape (seq_len, hidden_size).")
236

237
238
239
240
241
        # 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()

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

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

        max_length = tokenization_kwargs["max_length"]

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

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

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

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

        return inputs

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

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

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

        return inputs
320

321
    @overload
322
    def _prompt_to_llm_inputs(
323
        self,
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
350
        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,
351
        tokenization_kwargs: dict[str, Any] | None = None,
352
        *,
353
        mm_uuids: MultiModalUUIDDict | None = None,
354
    ) -> SingletonInputs:
355
356
        """
        Extract the singleton inputs from a prompt.
357
358
359

        Arguments:

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

        Returns:

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

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

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

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

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

390
391
392
393
        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."
394
            )
395

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

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

404
        return inputs
405

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

413
414
415
416
417
418
419
        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
420

421
422
423
        if enc_inputs["type"] == "multimodal":
            enc_inputs_new = token_inputs(enc_inputs["encoder_prompt_token_ids"])
            dec_inputs_new = MultiModalInputs(
424
                type="multimodal",
425
426
427
428
                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"],
429
            )
430
431
432
        elif enc_inputs["type"] == "token":
            enc_inputs_new = token_inputs(prompt_token_ids=[])
            dec_inputs_new = dec_inputs
433
        else:
434
435
436
437
438
439
440
            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
441

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

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

        Arguments:

459
        * prompt: an input prompt
460
461
462

        Returns:

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

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

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

        Arguments:

499
        * prompt: input prompt
500
501
502

        Returns:

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

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

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

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

        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

551
    def stat_mm_cache(self) -> MultiModalCacheStats | None:
552
553
554
555
556
557
558
559
560
        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:
561
562
        if self.mm_processor_cache is not None:
            self.mm_processor_cache.clear_cache()
563
564
565

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