cli.py 7.05 KB
Newer Older
Olivier Dehaene's avatar
v0.1.0  
Olivier Dehaene committed
1
import os
2
import sys
Olivier Dehaene's avatar
Olivier Dehaene committed
3
4
5
import typer

from pathlib import Path
6
from loguru import logger
7
from typing import Optional
8
from enum import Enum
9
from huggingface_hub import hf_hub_download
Olivier Dehaene's avatar
Olivier Dehaene committed
10
11
12
13
14


app = typer.Typer()


15
16
17
18
19
class Quantization(str, Enum):
    bitsandbytes = "bitsandbytes"
    gptq = "gptq"


20
21
22
23
24
class Dtype(str, Enum):
    float16 = "float16"
    bloat16 = "bfloat16"


Olivier Dehaene's avatar
Olivier Dehaene committed
25
@app.command()
Olivier Dehaene's avatar
v0.1.0  
Olivier Dehaene committed
26
def serve(
27
    model_id: str,
28
    revision: Optional[str] = None,
Olivier Dehaene's avatar
v0.1.0  
Olivier Dehaene committed
29
    sharded: bool = False,
30
    quantize: Optional[Quantization] = None,
31
    dtype: Optional[Dtype] = None,
32
    trust_remote_code: bool = False,
33
    uds_path: Path = "/tmp/text-generation-server",
34
35
    logger_level: str = "INFO",
    json_output: bool = False,
36
    otlp_endpoint: Optional[str] = None,
Olivier Dehaene's avatar
Olivier Dehaene committed
37
):
Olivier Dehaene's avatar
v0.1.0  
Olivier Dehaene committed
38
39
40
41
42
43
44
45
46
47
48
49
50
51
    if sharded:
        assert (
            os.getenv("RANK", None) is not None
        ), "RANK must be set when sharded is True"
        assert (
            os.getenv("WORLD_SIZE", None) is not None
        ), "WORLD_SIZE must be set when sharded is True"
        assert (
            os.getenv("MASTER_ADDR", None) is not None
        ), "MASTER_ADDR must be set when sharded is True"
        assert (
            os.getenv("MASTER_PORT", None) is not None
        ), "MASTER_PORT must be set when sharded is True"

52
53
54
55
56
    # Remove default handler
    logger.remove()
    logger.add(
        sys.stdout,
        format="{message}",
57
        filter="text_generation_server",
58
59
60
61
62
        level=logger_level,
        serialize=json_output,
        backtrace=True,
        diagnose=False,
    )
63
64
65
66
67

    # Import here after the logger is added to log potential import exceptions
    from text_generation_server import server
    from text_generation_server.tracing import setup_tracing

68
69
70
71
    # Setup OpenTelemetry distributed tracing
    if otlp_endpoint is not None:
        setup_tracing(shard=os.getenv("RANK", 0), otlp_endpoint=otlp_endpoint)

72
73
    # Downgrade enum into str for easier management later on
    quantize = None if quantize is None else quantize.value
74
75
76
77
78
79
80
81
    dtype = None if dtype is None else dtype.value
    if dtype is not None and quantize is not None:
        raise RuntimeError(
            "Only 1 can be set between `dtype` and `quantize`, as they both decide how goes the final model."
        )
    server.serve(
        model_id, revision, sharded, quantize, dtype, trust_remote_code, uds_path
    )
Olivier Dehaene's avatar
Olivier Dehaene committed
82
83
84


@app.command()
Nicolas Patry's avatar
Nicolas Patry committed
85
def download_weights(
86
    model_id: str,
87
    revision: Optional[str] = None,
88
    extension: str = ".safetensors",
89
    auto_convert: bool = True,
90
91
    logger_level: str = "INFO",
    json_output: bool = False,
92
    trust_remote_code: bool = False,
Olivier Dehaene's avatar
Olivier Dehaene committed
93
):
94
95
96
97
98
    # Remove default handler
    logger.remove()
    logger.add(
        sys.stdout,
        format="{message}",
99
        filter="text_generation_server",
100
101
102
103
104
105
        level=logger_level,
        serialize=json_output,
        backtrace=True,
        diagnose=False,
    )

106
107
108
    # Import here after the logger is added to log potential import exceptions
    from text_generation_server import utils

109
110
111
    # Test if files were already download
    try:
        utils.weight_files(model_id, revision, extension)
112
        logger.info("Files are already present on the host. " "Skipping download.")
113
114
        return
    # Local files not found
115
    except (utils.LocalEntryNotFoundError, FileNotFoundError):
116
117
        pass

118
119
120
121
122
    is_local_model = (Path(model_id).exists() and Path(model_id).is_dir()) or os.getenv(
        "WEIGHTS_CACHE_OVERRIDE", None
    ) is not None

    if not is_local_model:
123
124
125
126
127
128
        try:
            adapter_config_filename = hf_hub_download(model_id, revision=revision, filename="adapter_config.json")
            utils.download_and_unload_peft(model_id, revision, trust_remote_code=trust_remote_code)
        except (utils.LocalEntryNotFoundError, utils.EntryNotFoundError):
            pass

129
130
131
132
133
134
135
136
137
138
139
140
141
142
        # Try to download weights from the hub
        try:
            filenames = utils.weight_hub_files(model_id, revision, extension)
            utils.download_weights(filenames, model_id, revision)
            # Successfully downloaded weights
            return

        # No weights found on the hub with this extension
        except utils.EntryNotFoundError as e:
            # Check if we want to automatically convert to safetensors or if we can use .bin weights instead
            if not extension == ".safetensors" or not auto_convert:
                raise e

    # Try to see if there are local pytorch weights
143
    try:
144
145
        # Get weights for a local model, a hub cached model and inside the WEIGHTS_CACHE_OVERRIDE
        local_pt_files = utils.weight_files(model_id, revision, ".bin")
146

147
148
149
150
151
152
153
    # No local pytorch weights
    except utils.LocalEntryNotFoundError:
        if extension == ".safetensors":
            logger.warning(
                f"No safetensors weights found for model {model_id} at revision {revision}. "
                f"Downloading PyTorch weights."
            )
154

155
        # Try to see if there are pytorch weights on the hub
156
157
158
        pt_filenames = utils.weight_hub_files(model_id, revision, ".bin")
        # Download pytorch weights
        local_pt_files = utils.download_weights(pt_filenames, model_id, revision)
159
160
161
162
163
164
165
166

    if auto_convert:
        logger.warning(
            f"No safetensors weights found for model {model_id} at revision {revision}. "
            f"Converting PyTorch weights to safetensors."
        )

        # Safetensors final filenames
167
168
169
170
        local_st_files = [
            p.parent / f"{p.stem.lstrip('pytorch_')}.safetensors"
            for p in local_pt_files
        ]
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
        try:
            from transformers import AutoConfig
            import transformers

            config = AutoConfig.from_pretrained(
                model_id,
                revision=revision,
            )
            architecture = config.architectures[0]

            class_ = getattr(transformers, architecture)

            # Name for this varible depends on transformers version.
            discard_names = getattr(class_, "_tied_weights_keys", [])
            discard_names.extend(getattr(class_, "_keys_to_ignore_on_load_missing", []))

        except Exception as e:
            discard_names = []
189
        # Convert pytorch weights to safetensors
190
        utils.convert_files(local_pt_files, local_st_files, discard_names)
Olivier Dehaene's avatar
Olivier Dehaene committed
191
192


193
194
195
196
197
198
199
200
201
202
203
204
@app.command()
def quantize(
    model_id: str,
    output_dir: str,
    revision: Optional[str] = None,
    logger_level: str = "INFO",
    json_output: bool = False,
    trust_remote_code: bool = False,
    upload_to_model_id: Optional[str] = None,
    percdamp: float = 0.01,
    act_order: bool = False,
):
205
206
    if revision is None:
        revision = "main"
207
208
209
210
211
212
213
214
215
216
217
218
219
    download_weights(
        model_id=model_id,
        revision=revision,
        logger_level=logger_level,
        json_output=json_output,
    )
    from text_generation_server.utils.gptq.quantize import quantize

    quantize(
        model_id=model_id,
        bits=4,
        groupsize=128,
        output_dir=output_dir,
220
        revision=revision,
221
222
223
224
225
226
227
        trust_remote_code=trust_remote_code,
        upload_to_model_id=upload_to_model_id,
        percdamp=percdamp,
        act_order=act_order,
    )


Olivier Dehaene's avatar
Olivier Dehaene committed
228
229
if __name__ == "__main__":
    app()