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

4
from abc import ABC, abstractmethod
5
from collections.abc import AsyncGenerator, Iterable, Mapping
6
from typing import Any
7

8
from vllm.config import ModelConfig, VllmConfig
9
from vllm.inputs.data import PromptType
10
from vllm.lora.request import LoRARequest
11
12
from vllm.outputs import PoolingRequestOutput, RequestOutput
from vllm.plugins.io_processors import IOProcessor
13
from vllm.pooling_params import PoolingParams
14
from vllm.sampling_params import SamplingParams
15
from vllm.tasks import SupportedTask
16
from vllm.tokenizers import TokenizerLike
17
from vllm.v1.engine import EngineCoreRequest
18
from vllm.v1.engine.input_processor import InputProcessor
19

20

21
class EngineClient(ABC):
22
    """Protocol class for Clients to Engine"""
23

24
25
    vllm_config: VllmConfig
    model_config: ModelConfig
26
    input_processor: InputProcessor
27
    io_processor: IOProcessor | None
28

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

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

    @property
38
    @abstractmethod
39
    def errored(self) -> bool: ...
40

41
    @property
42
    @abstractmethod
43
    def dead_error(self) -> BaseException: ...
44

45
    @abstractmethod
46
    def generate(
47
        self,
48
        prompt: EngineCoreRequest | PromptType,
49
50
        sampling_params: SamplingParams,
        request_id: str,
51
        *,
52
53
54
55
        prompt_text: str | None = None,
        lora_request: LoRARequest | None = None,
        tokenization_kwargs: dict[str, Any] | None = None,
        trace_headers: Mapping[str, str] | None = None,
56
        priority: int = 0,
57
        data_parallel_rank: int | None = None,
58
    ) -> AsyncGenerator[RequestOutput, None]:
59
        """Generate outputs for a request."""
60
        ...
61

62
    @abstractmethod
63
    def encode(
64
        self,
65
        prompt: PromptType,
66
67
        pooling_params: PoolingParams,
        request_id: str,
68
69
        lora_request: LoRARequest | None = None,
        trace_headers: Mapping[str, str] | None = None,
70
        priority: int = 0,
71
        truncate_prompt_tokens: int | None = None,
72
        tokenization_kwargs: dict[str, Any] | None = None,
73
    ) -> AsyncGenerator[PoolingRequestOutput, None]:
74
75
76
77
78
        """Generate outputs for a request from a pooling model.

        NOTE: truncate_prompt_tokens is deprecated in v0.14.
        TODO: Remove this argument in v0.15.
        """
79
        ...
80

81
    @abstractmethod
82
    async def abort(self, request_id: str | Iterable[str]) -> None:
83
84
85
        """Abort a request.

        Args:
86
87
            request_id: The unique id of the request,
                        or an iterable of such ids.
88
        """
89
        ...
90

91
    @abstractmethod
92
    async def get_tokenizer(self) -> TokenizerLike:
93
        """Get the tokenizer"""
94
        ...
95

96
    @abstractmethod
97
    async def is_tracing_enabled(self) -> bool: ...
98

99
    @abstractmethod
100
    async def do_log_stats(self) -> None: ...
101

102
    @abstractmethod
103
104
    async def check_health(self) -> None:
        """Raise if unhealthy"""
105
        ...
106

107
    @abstractmethod
108
109
110
111
    async def start_profile(self) -> None:
        """Start profiling the engine"""
        ...

112
    @abstractmethod
113
    async def stop_profile(self) -> None:
114
        """Stop profiling the engine"""
115
        ...
116

117
118
119
120
121
    @abstractmethod
    async def reset_mm_cache(self) -> None:
        """Reset the multi-modal cache"""
        ...

122
    @abstractmethod
123
124
125
126
    async def reset_prefix_cache(
        self, reset_running_requests: bool = False, reset_connector: bool = False
    ) -> bool:
        """Reset the prefix cache and optionally any configured connector cache"""
127
128
        ...

129
130
131
132
133
134
    @abstractmethod
    async def sleep(self, level: int = 1) -> None:
        """Sleep the engine"""
        ...

    @abstractmethod
135
    async def wake_up(self, tags: list[str] | None = None) -> None:
136
137
138
        """Wake up the engine"""
        ...

139
140
141
142
143
    @abstractmethod
    async def is_sleeping(self) -> bool:
        """Check whether the engine is sleeping"""
        ...

144
    @abstractmethod
145
    async def add_lora(self, lora_request: LoRARequest) -> bool:
146
147
        """Load a new LoRA adapter into the engine for future requests."""
        ...
148

149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
    @abstractmethod
    async def pause_generation(
        self,
        *,
        wait_for_inflight_requests: bool = False,
        clear_cache: bool = True,
    ) -> None:
        """Pause new generation/encoding requests.

        Args:
            wait_for_inflight_requests: When ``True`` waits for in-flight requests
                to finish before pausing. When ``False`` (default), aborts in-flight
                requests immediately.
            clear_cache: Whether to clear KV and prefix caches after draining.
        """
        ...

    @abstractmethod
    async def resume_generation(self) -> None:
        """Resume accepting generation/encoding requests."""
        ...

    @abstractmethod
    async def is_paused(self) -> bool:
        """Return whether the engine is currently paused."""
        ...

176
177
178
    async def scale_elastic_ep(
        self, new_data_parallel_size: int, drain_timeout: int = 300
    ) -> None:
179
180
        """Scale the engine"""
        raise NotImplementedError
181

182
183
184
    async def collective_rpc(
        self,
        method: str,
185
        timeout: float | None = None,
186
        args: tuple = (),
187
        kwargs: dict | None = None,
188
    ):
189
190
        """Perform a collective RPC call to the given path."""
        raise NotImplementedError
191
192
193
194

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