E: Could not open lock file /var/lib/dpkg/lock-frontend - open (13: Permission denied)
E: Unable to acquire the dpkg frontend lock (/var/lib/dpkg/lock-frontend), are you root?
2022-12-14 21:23:54.920413: W tensorflow/compiler/xla/stream_executor/platform/default/dso_loader.cc:64] Could not load dynamic library 'libnvinfer.so.7'; dlerror: libnvinfer.so.7: cannot open shared object file: No such file or directory
2022-12-14 21:23:54.920526: W tensorflow/compiler/xla/stream_executor/platform/default/dso_loader.cc:64] Could not load dynamic library 'libnvinfer_plugin.so.7'; dlerror: libnvinfer_plugin.so.7: cannot open shared object file: No such file or directory
2022-12-14 21:23:54.920537: W tensorflow/compiler/tf2tensorrt/utils/py_utils.cc:38] TF-TRT Warning: Cannot dlopen some TensorRT libraries. If you would like to use Nvidia GPU with TensorRT, please make sure the missing libraries mentioned above are installed properly.
[선택 사항] 데이터 처리
이 섹션은 캡션 데이터세트를 다운로드하고 훈련을 위해 이를 준비합니다. 입력 텍스트를 토큰화하고 사전 훈련된 특정 추출 모델을 통해 모든 이미지를 실행한 결과를 캐싱합니다. 이는 이 섹션의 모든 것을 이해하는 데 중요하지는 않습니다.
defconceptual_captions(*,data_dir="conceptual_captions",num_train,num_val):defiter_index(index_path):withopen(index_path)asf:forlineinf:caption,url=line.strip().split('\t')yieldcaption,urldefdownload_image_urls(data_dir,urls):ex=concurrent.futures.ThreadPoolExecutor(max_workers=100)defsave_image(url):hash=hashlib.sha1(url.encode())# Name the files after the hash of the URL.file_path=data_dir/f'{hash.hexdigest()}.jpeg'iffile_path.exists():# Only download each file once.returnfile_pathtry:result=requests.get(url,timeout=5)exceptException:file_path=Noneelse:file_path.write_bytes(result.content)returnfile_pathresult=[]out_paths=ex.map(save_image,urls)forfile_pathintqdm.tqdm(out_paths,total=len(urls)):result.append(file_path)returnresultdefds_from_index_file(index_path,data_dir,count):data_dir.mkdir(exist_ok=True)index=list(itertools.islice(iter_index(index_path),count))captions=[captionforcaption,urlinindex]urls=[urlforcaption,urlinindex]paths=download_image_urls(data_dir,urls)new_captions=[]new_paths=[]forcap,pathinzip(captions,paths):ifpathisNone:# Download failed, so skip this pair.continuenew_captions.append(cap)new_paths.append(path)new_paths=[str(p)forpinnew_paths]ds=tf.data.Dataset.from_tensor_slices((new_paths,new_captions))ds=ds.map(lambdapath,cap:(path,cap[tf.newaxis]))# 1 caption per imagereturndsdata_dir=pathlib.Path(data_dir)train_index_path=tf.keras.utils.get_file(origin='https://storage.googleapis.com/gcc-data/Train/GCC-training.tsv',cache_subdir=data_dir,cache_dir='.')val_index_path=tf.keras.utils.get_file(origin='https://storage.googleapis.com/gcc-data/Validation/GCC-1.1.0-Validation.tsv',cache_subdir=data_dir,cache_dir='.')train_raw=ds_from_index_file(train_index_path,data_dir=data_dir/'train',count=num_train)test_raw=ds_from_index_file(val_index_path,data_dir=data_dir/'val',count=num_val)returntrain_raw,test_raw
데이터세트 다운로드
Flickr8k는 이미지당 5개의 캡션과 더욱 소규모의 다운로드를 위한 더 많은 데이터를 포함하고 있어 좋은 선택입니다.
Downloading data from https://github.com/jbrownlee/Datasets/releases/download/Flickr8k/Flickr8k_Dataset.zip
1115419746/1115419746 [==============================] - 25s 0us/step
Downloading data from https://github.com/jbrownlee/Datasets/releases/download/Flickr8k/Flickr8k_text.zip
2340801/2340801 [==============================] - 1s 0us/step
위의 두 데이터세트에 대한 로더는 (image_path, captions) 쌍을 포함하는 tf.data.Dataset를 반환합니다. Conceptual Captions는 이미지당 캡션 1개를 포함하는 한편 Flickr8k는 이미지당 5개의 캡션을 포함합니다.
tf.Tensor(b'flickr8k/Flicker8k_Dataset/2513260012_03d33305cf.jpg', shape=(), dtype=string)
tf.Tensor(
[b'A black dog is running after a white dog in the snow .'
b'Black dog chasing brown dog through snow'
b'Two dogs chase each other across the snowy ground .'
b'Two dogs play together in the snow .'
b'Two dogs running through a low lying body of water .'], shape=(5,), dtype=string)
이미지 특성 추출기
각 이미지에서 특성을 추출하기 위해 이미지 모델(imagenet에서 사전 훈련됨)을 사용할 것입니다. 모델은 이미지 분류기로 훈련되었지만, 설정 include_top=False는 최종 분류 레이어 없이 모델을 반환하므로 특성 맵의 최종 레이어를 사용할 수 있습니다.
Downloading data from https://storage.googleapis.com/tensorflow/keras-applications/mobilenet_v3/weights_mobilenet_v3_small_224_1.0_float_no_top_v2.h5
4334752/4334752 [==============================] - 0s 0us/step
# Use the top 5000 words for a vocabulary.vocabulary_size=5000tokenizer=tf.keras.layers.TextVectorization(max_tokens=vocabulary_size,standardize=standardize,ragged=True)# Learn the vocabulary from the caption data.
WARNING:tensorflow:From /tmpfs/src/tf_docs_env/lib/python3.9/site-packages/tensorflow/python/autograph/pyct/static_analysis/liveness.py:83: Analyzer.lamba_check (from tensorflow.python.autograph.pyct.static_analysis.liveness) is deprecated and will be removed after 2023-09-23.
Instructions for updating:
Lambda fuctions will be no more assumed to be used in the statement where they are used, or at least in the same block. https://github.com/tensorflow/tensorflow/issues/56089
# Create mappings for words to indices and indices to words.word_to_index=tf.keras.layers.StringLookup(mask_token="",vocabulary=tokenizer.get_vocabulary())index_to_word=tf.keras.layers.StringLookup(mask_token="",vocabulary=tokenizer.get_vocabulary(),invert=True)
keras 훈련과 호환되려면 데이터세트는 (inputs, labels) 쌍을 포함해야 합니다. 텍스트 생성의 경우 토큰은 한 단계 이동된 입력과 라벨입니다. 이 함수는 (images, texts) 쌍을 ((images, input_tokens), label_tokens) 쌍으로 변환합니다.
RaggedTensor 표현에서 텍스트를 패딩 처리된 밀도 높은 Tensor 표현으로 변환합니다.
defprepare_dataset(ds,tokenizer,batch_size=32,shuffle_buffer=1000):# Load the images and make batches.ds=(ds.shuffle(10000).map(lambdapath,caption:(load_image(path),caption)).apply(tf.data.experimental.ignore_errors()).batch(batch_size))defto_tensor(inputs,labels):(images,in_tok),out_tok=inputs,labelsreturn(images,in_tok.to_tensor()),out_tok.to_tensor()return(ds.map(match_shapes,tf.data.AUTOTUNE).unbatch().shuffle(shuffle_buffer).batch(batch_size).map(prepare_txt,tf.data.AUTOTUNE).map(to_tensor,tf.data.AUTOTUNE))
WARNING:tensorflow:From /tmpfs/tmp/ipykernel_676096/1004139779.py:6: ignore_errors (from tensorflow.python.data.experimental.ops.error_ops) is deprecated and will be removed in a future version.
Instructions for updating:
Use `tf.data.Dataset.ignore_errors` instead.
((TensorSpec(shape=(None, 224, 224, 3), dtype=tf.float32, name=None),
TensorSpec(shape=(None, None), dtype=tf.int64, name=None)),
TensorSpec(shape=(None, None), dtype=tf.int64, name=None))
이미지 특성 추출기가 변경되지 않으며 이 튜토리얼은 이미지 증강을 사용하지 않으므로 이미지 특성은 캐싱될 수 있습니다. 텍스트 토큰화의 경우도 동일합니다. 캐시를 설정하는 데 드는 시간은 훈련 및 검증 중 각 epoch에서 다시 획득됩니다. 아래의 코드는 두 개의 함수인 save_dataset 및 load_dataset를 정의합니다.
defsave_dataset(ds,save_path,image_model,tokenizer,shards=10,batch_size=32):# Load the images and make batches.ds=(ds.map(lambdapath,caption:(load_image(path),caption)).apply(tf.data.experimental.ignore_errors()).batch(batch_size))# Run the feature extractor on each batch# Don't do this in a .map, because tf.data runs on the CPU. defgen():for(images,captions)intqdm.tqdm(ds):feature_maps=image_model(images)feature_maps,captions=match_shapes(feature_maps,captions)yieldfeature_maps,captions# Wrap the generator in a new tf.data.Dataset.new_ds=tf.data.Dataset.from_generator(gen,output_signature=(tf.TensorSpec(shape=image_model.output_shape),tf.TensorSpec(shape=(None,),dtype=tf.string)))# Apply the tokenization new_ds=(new_ds.map(prepare_txt,tf.data.AUTOTUNE).unbatch().shuffle(1000))# Save the dataset into shard files.defshard_func(i,item):returni%shardsnew_ds.enumerate().save(save_path,shard_func=shard_func)defload_dataset(save_path,batch_size=32,shuffle=1000,cycle_length=2):defcustom_reader_func(datasets):datasets=datasets.shuffle(1000)returndatasets.interleave(lambdax:x,cycle_length=cycle_length)ds=tf.data.Dataset.load(save_path,reader_func=custom_reader_func)defdrop_index(i,x):returnxds=(ds.map(drop_index,tf.data.AUTOTUNE).shuffle(shuffle).padded_batch(batch_size).prefetch(tf.data.AUTOTUNE))returnds
데이터세트는 이제 keras 훈련에 적합한 (input, label) 쌍을 반환합니다. inputs은 (images, input_tokens) 쌍입니다. images는 특성-추출기 모델로 처리됩니다. input_tokens의 각 위치의 경우 모델은 지금까지의 텍스트를 보고 labels의 같은 위치에서 나열된 다음 텍스트를 예측하려고 시도합니다.
추후에 각 출력 위치가 지금까지 출력에 대해 처리할 수 있는 인과적 셀프 어텐션(CausalSelfAttention).
각 출력 위치가 입력 이미지를 추리할 수 있는 크로스 어텐션 레이어(CrossAttention).
각 출력 위치를 독립적으로 추가로 처리하는 피드 포워드 네트워크(FeedForward) 레이어.
출력 - 출력 어휘에 대한 멀티 클래스 분류.
입력
입력 텍스트는 이미 토큰으로 분할되고 ID 시퀀스로 변환되었습니다.
CNN 또는 RNN와는 다르게 트랜스포머의 어텐션 레이어는 시퀀스의 순서에 대해 변하지 않는다는 점을 기억하세요. 몇몇 위치 입력이 없다면 시퀀스가 아닌 순서 없는 세트만 봅니다. 따라서 각 토큰 ID에 대한 단순한 벡터 임베딩 외에도 임베딩 레이어는 시퀀스 내 각 위치에 대한 임베딩 또한 포함합니다.
SeqEmbedding 레이어는 다음과 같이 정의됩니다.
각 토큰에 대한 임베딩 벡터를 검색합니다.
각 시퀀스 위치에 대한 임베딩 벡터를 검색합니다.
두 개를 모두 합합니다.
mask_zero=True를 사용하여 모델에 대한 keras 마스크를 초기화합니다.
참고: 이 구현은 Transformer 튜토리얼에서와 같이 고정된 임베딩을 사용하는 대신 위치 임베딩을 학습합니다. 임베딩을 학습하는 것은 코드가 약간 적지만 더 긴 시퀀스로 일반화되지는 않습니다.
디코더는 표준 트랜스포머 디코더로, 각 세 개의 하위 레이어인 CausalSelfAttention, CrossAttention 및 FeedForward를 포함하는 DecoderLayers의 스택을 포함합니다. 구현은 Transformer 튜토리얼과 거의 동일하며, 자세한 내용은 이를 참조하세요.
다음은 CausalSelfAttention 레이어입니다.
classCausalSelfAttention(tf.keras.layers.Layer):def__init__(self,**kwargs):super().__init__()self.mha=tf.keras.layers.MultiHeadAttention(**kwargs)# Use Add instead of + so the keras mask propagates through.self.add=tf.keras.layers.Add()self.layernorm=tf.keras.layers.LayerNormalization()defcall(self,x):attn=self.mha(query=x,value=x,use_causal_mask=True)x=self.add([x,attn])returnself.layernorm(x)
아래는 CrossAttention 레이어입니다. return_attention_scores를 사용하는 데 유의하세요.
다음으로 이러한 세 가지 레이어를 더 큰 규모의 DecoderLayer에 배열합니다. 각 디코더 레이어는 시퀀스에 세 개의 더 작은 레이어를 적용합니다. 각 하위 레이어 다음의 out_seq 형태는 (batch, sequence, channels)입니다. 디코더 레이어는 또한 추후 시각화를 위한 attention_scores를 반환합니다.
classDecoderLayer(tf.keras.layers.Layer):def__init__(self,units,num_heads=1,dropout_rate=0.1):super().__init__()self.self_attention=CausalSelfAttention(num_heads=num_heads,key_dim=units,dropout=dropout_rate)self.cross_attention=CrossAttention(num_heads=num_heads,key_dim=units,dropout=dropout_rate)self.ff=FeedForward(units=units,dropout_rate=dropout_rate)defcall(self,inputs,training=False):in_seq,out_seq=inputs# Text inputout_seq=self.self_attention(out_seq)out_seq=self.cross_attention(out_seq,in_seq)self.last_attention_scores=self.cross_attention.last_attention_scoresout_seq=self.ff(out_seq)returnout_seq
출력
출력 레이어는 각 위치에서 각 토큰에 대한 로짓 예측을 생성하려면 최소한 layers.Dense 레이어가 필요합니다.
하지만 이 작업을 좀 더 잘 수행할 수 있도록 추가할 수 있는 몇 가지 다른 특성이 있습니다.
잘못된 토큰 처리: 모델은 텍스트를 생성합니다. 패드, 알 수 없는, 또는 시작 토큰('', '[UNK]', '[START]')을 생성해서는 안됩니다. 따라서 이들에 대한 편향을 큰 음수 값으로 설정합니다.
참고: 손실 함수의 이러한 토큰 역시 무시해야 합니다.
스마트 초기화: 밀도가 높은 레이어의 기본 초기화는 거의 균일한 확률로 각 토큰을 초기에 예측하는 모델을 제공합니다. 실제 토큰 분포는 균일한 것과는 거리가 멉니다. 출력 레이어의 초기 편향을 위한 최적값은 각 토큰의 확률 로그입니다. 따라서 adapt 메서드를 포함해 토큰의 수를 세고 최적의 초기 편향을 설정합니다. 이는 균일한 분포(log(vocabulary_size))의 엔트로피로부터의 분포의 한계 엔트로피(-p*log(p))로 초기 손실을 줄입니다.
classTokenOutput(tf.keras.layers.Layer):def__init__(self,tokenizer,banned_tokens=('','[UNK]','[START]'),**kwargs):super().__init__()self.dense=tf.keras.layers.Dense(units=tokenizer.vocabulary_size(),**kwargs)self.tokenizer=tokenizerself.banned_tokens=banned_tokensself.bias=Nonedefadapt(self,ds):counts=collections.Counter()vocab_dict={name:idforid,nameinenumerate(self.tokenizer.get_vocabulary())}fortokensintqdm.tqdm(ds):counts.update(tokens.numpy().flatten())counts_arr=np.zeros(shape=(self.tokenizer.vocabulary_size(),))counts_arr[np.array(list(counts.keys()),dtype=np.int32)]=list(counts.values())counts_arr=counts_arr[:]fortokeninself.banned_tokens:counts_arr[vocab_dict[token]]=0total=counts_arr.sum()p=counts_arr/totalp[counts_arr==0]=1.0log_p=np.log(p)# log(1) == 0entropy=-(log_p*p).sum()print()print(f"Uniform entropy: {np.log(self.tokenizer.vocabulary_size()):0.2f}")print(f"Marginal entropy: {entropy:0.2f}")self.bias=log_pself.bias[counts_arr==0]=-1e9defcall(self,x):x=self.dense(x)# TODO(b/250038731): Fix this.# An Add layer doesn't work because of the different shapes.# This clears the mask, that's okay because it prevents keras from rescaling# the losses.returnx+self.bias
스마트 초기화는 초기 손실을 다음과 같이 상당히 줄입니다.
output_layer=TokenOutput(tokenizer,banned_tokens=('','[UNK]','[START]'))# This might run a little faster if the dataset didn't also have to load the image data.output_layer.adapt(train_ds.map(lambdainputs,labels:labels))
@Captioner.add_methoddefcall(self,inputs):image,txt=inputsifimage.shape[-1]==3:# Apply the feature-extractor, if you get an RGB image.image=self.feature_extractor(image)# Flatten the feature mapimage=einops.rearrange(image,'b h w c -> b (h w) c')iftxt.dtype==tf.string:# Apply the tokenizer if you get string inputs.txt=tokenizer(txt)txt=self.seq_embedding(txt)# Look at the imagefordec_layerinself.decoder_layers:txt=dec_layer(inputs=(image,txt))txt=self.output_layer(txt)returntxt
Epoch 1/100
99/100 [============================>.] - ETA: 0s - loss: 4.9987 - masked_acc: 0.2017
a man in a man in a man
a man in the white in the running
a inflatable child parlor crowd in top grinds
100/100 [==============================] - 24s 140ms/step - loss: 4.9973 - masked_acc: 0.2018 - val_loss: 4.6630 - val_masked_acc: 0.2443
Epoch 2/100
100/100 [==============================] - ETA: 0s - loss: 4.6423 - masked_acc: 0.2546
a man in a black dog is in a man
a woman in a man of a white white a white yellow in a red in a black
a woman her air with black boy run his is players of an another to a towards rides a stand a
100/100 [==============================] - 9s 87ms/step - loss: 4.6423 - masked_acc: 0.2546 - val_loss: 4.3656 - val_masked_acc: 0.2699
Epoch 3/100
98/100 [============================>.] - ETA: 0s - loss: 4.4123 - masked_acc: 0.2763
a man in a red in the water
a boy is and a small boy in a brown child
two two at a green down a water in the hanging
100/100 [==============================] - 7s 66ms/step - loss: 4.4109 - masked_acc: 0.2766 - val_loss: 4.1544 - val_masked_acc: 0.2931
Epoch 4/100
98/100 [============================>.] - ETA: 0s - loss: 4.2374 - masked_acc: 0.2961
a man in a red shirt is in a red and a red
a boy in a yellow through the water
a motorcyclist a in front of water
100/100 [==============================] - 6s 57ms/step - loss: 4.2346 - masked_acc: 0.2967 - val_loss: 4.0267 - val_masked_acc: 0.3058
Epoch 5/100
100/100 [==============================] - ETA: 0s - loss: 4.1222 - masked_acc: 0.3088
a man in a red shirt is in the water
a man is in the water
a woman down the the
100/100 [==============================] - 5s 50ms/step - loss: 4.1222 - masked_acc: 0.3088 - val_loss: 3.9308 - val_masked_acc: 0.3205
Epoch 6/100
100/100 [==============================] - ETA: 0s - loss: 4.0236 - masked_acc: 0.3136
a man in a red shirt is in the water
a man is sitting in the water
a blue team is playing on the water
100/100 [==============================] - 5s 54ms/step - loss: 4.0236 - masked_acc: 0.3136 - val_loss: 3.8872 - val_masked_acc: 0.3231
Epoch 7/100
100/100 [==============================] - ETA: 0s - loss: 3.9255 - masked_acc: 0.3239
a man in a blue shirt is jumping in the water
a man in a large water
the on the canoe snowy swinging
100/100 [==============================] - 5s 50ms/step - loss: 3.9255 - masked_acc: 0.3239 - val_loss: 3.8013 - val_masked_acc: 0.3318
Epoch 8/100
98/100 [============================>.] - ETA: 0s - loss: 3.8617 - masked_acc: 0.3311
a man in a blue shirt is jumping in the water
a woman in a red shirt is in the water
a man on a a holds water on a picture of the water
100/100 [==============================] - 6s 57ms/step - loss: 3.8677 - masked_acc: 0.3308 - val_loss: 3.7248 - val_masked_acc: 0.3376
Epoch 9/100
99/100 [============================>.] - ETA: 0s - loss: 3.8062 - masked_acc: 0.3335
a man in a blue shirt is jumping in the water
a man is a yellow shirt is walking through a pool
a money a and red shirt in sing in the wave
100/100 [==============================] - 6s 60ms/step - loss: 3.8032 - masked_acc: 0.3338 - val_loss: 3.6805 - val_masked_acc: 0.3426
Epoch 10/100
99/100 [============================>.] - ETA: 0s - loss: 3.7149 - masked_acc: 0.3418
a man in a red shirt is jumping in the water
a man in a red shirt is jumping in a red is holding a blue above the water
a boy girl in leaves above the ocean
100/100 [==============================] - 6s 60ms/step - loss: 3.7143 - masked_acc: 0.3415 - val_loss: 3.6240 - val_masked_acc: 0.3424
Epoch 11/100
99/100 [============================>.] - ETA: 0s - loss: 3.6224 - masked_acc: 0.3491
a man in a blue shirt is jumping into the water
person on a water in the water
a woman playing a dock off some waterfall
100/100 [==============================] - 5s 53ms/step - loss: 3.6225 - masked_acc: 0.3489 - val_loss: 3.5840 - val_masked_acc: 0.3431
Epoch 12/100
99/100 [============================>.] - ETA: 0s - loss: 3.5929 - masked_acc: 0.3519
a man in a blue shirt is jumping in the water
a man in a blue shirt is jumping over a pool
a small boy in beer in a river
100/100 [==============================] - 6s 60ms/step - loss: 3.5916 - masked_acc: 0.3518 - val_loss: 3.4882 - val_masked_acc: 0.3528
Epoch 13/100
100/100 [==============================] - ETA: 0s - loss: 3.5806 - masked_acc: 0.3484
a man in a red shirt is swimming pool
a person and a red dog is swimming pool
two women pose to a skate in a deck
100/100 [==============================] - 5s 52ms/step - loss: 3.5806 - masked_acc: 0.3484 - val_loss: 3.4243 - val_masked_acc: 0.3611
Epoch 14/100
99/100 [============================>.] - ETA: 0s - loss: 3.5164 - masked_acc: 0.3566
a man in a red shirt is jumping into the water
a man in a blue shirt is looking at a wave
a skateboarder is performing in the ocean
100/100 [==============================] - 5s 55ms/step - loss: 3.5139 - masked_acc: 0.3570 - val_loss: 3.4775 - val_masked_acc: 0.3508
Epoch 15/100
99/100 [============================>.] - ETA: 0s - loss: 3.5156 - masked_acc: 0.3557
a man in a blue shirt is jumping into the water
a young boy in a wave
the person and a boy shoveling into the water
100/100 [==============================] - 5s 53ms/step - loss: 3.5120 - masked_acc: 0.3560 - val_loss: 3.4817 - val_masked_acc: 0.3532
Epoch 16/100
99/100 [============================>.] - ETA: 0s - loss: 3.4653 - masked_acc: 0.3603
a man in a red shirt is riding a wave
a man in a blue shirt is holding a blue water
a person leaps around a rough middle of a pool in his ice
100/100 [==============================] - 6s 58ms/step - loss: 3.4648 - masked_acc: 0.3603 - val_loss: 3.3895 - val_masked_acc: 0.3593
Epoch 17/100
100/100 [==============================] - ETA: 0s - loss: 3.4460 - masked_acc: 0.3639
a man in a red shirt is swimming pool
two people are in a pool
a yellow clothes is being in a boat with green ocean on a pool
100/100 [==============================] - 6s 58ms/step - loss: 3.4460 - masked_acc: 0.3639 - val_loss: 3.2745 - val_masked_acc: 0.3722
Epoch 18/100
99/100 [============================>.] - ETA: 0s - loss: 3.4242 - masked_acc: 0.3613
a man in a red shirt is swimming pool
a man in a blue shirt is swimming pool
a man in a white tshirt waves with his head in a orange high in the air
100/100 [==============================] - 6s 55ms/step - loss: 3.4241 - masked_acc: 0.3614 - val_loss: 3.2940 - val_masked_acc: 0.3669
Epoch 19/100
100/100 [==============================] - ETA: 0s - loss: 3.3747 - masked_acc: 0.3667
a man in a red shirt is jumping into the water
a man in a red hat and white jacket is jump into a wave
a greyhound a young woman standing in a pool
100/100 [==============================] - 6s 57ms/step - loss: 3.3747 - masked_acc: 0.3667 - val_loss: 3.2806 - val_masked_acc: 0.3624
Epoch 20/100
98/100 [============================>.] - ETA: 0s - loss: 3.2628 - masked_acc: 0.3771
a man in a red shirt is swimming pool
a man in a red shirt and white jacket is surfing
closeup kayak in an orange cast cyclist
100/100 [==============================] - 5s 51ms/step - loss: 3.2634 - masked_acc: 0.3765 - val_loss: 3.3122 - val_masked_acc: 0.3580
Epoch 21/100
100/100 [==============================] - ETA: 0s - loss: 3.2708 - masked_acc: 0.3758
a man in a red shirt is swimming pool
a boy is a blue wave on a wave
a person is mohawk is pose in a blue
100/100 [==============================] - 5s 52ms/step - loss: 3.2708 - masked_acc: 0.3758 - val_loss: 3.2734 - val_masked_acc: 0.3624
Epoch 22/100
99/100 [============================>.] - ETA: 0s - loss: 3.2577 - masked_acc: 0.3735
a man in a red shirt is swimming pool
a person in a blue shirt and white wave
a boy with a dog in a black tshirt in the water
100/100 [==============================] - 6s 58ms/step - loss: 3.2566 - masked_acc: 0.3737 - val_loss: 3.2527 - val_masked_acc: 0.3716
Epoch 23/100
99/100 [============================>.] - ETA: 0s - loss: 3.2335 - masked_acc: 0.3788
a man in a blue shirt is swimming pool
a boy wearing a blue shirt is in a pool
a surfer is surfing
100/100 [==============================] - 5s 50ms/step - loss: 3.2340 - masked_acc: 0.3788 - val_loss: 3.2174 - val_masked_acc: 0.3700
Epoch 24/100
98/100 [============================>.] - ETA: 0s - loss: 3.2065 - masked_acc: 0.3845
a man in a red shirt is riding a wave
a man in a red shirt is swimming pool
an old boy holds a wave picture at the a swimming pool
100/100 [==============================] - 6s 57ms/step - loss: 3.2100 - masked_acc: 0.3843 - val_loss: 3.1551 - val_masked_acc: 0.3777
Epoch 25/100
100/100 [==============================] - ETA: 0s - loss: 3.2055 - masked_acc: 0.3775
a man in a red and white wave
a man in red helmet rides a wave
a little child jumps onto the air in the ocean
100/100 [==============================] - 5s 53ms/step - loss: 3.2055 - masked_acc: 0.3775 - val_loss: 3.1546 - val_masked_acc: 0.3791
Epoch 26/100
99/100 [============================>.] - ETA: 0s - loss: 3.1624 - masked_acc: 0.3847
a man in a red shirt is swimming pool
a person in a red and yellow surfboard is riding a wave
the man is blowing snowmobile in shallow water
100/100 [==============================] - 5s 54ms/step - loss: 3.1614 - masked_acc: 0.3847 - val_loss: 3.1898 - val_masked_acc: 0.3706
Epoch 27/100
99/100 [============================>.] - ETA: 0s - loss: 3.1709 - masked_acc: 0.3834
a man in a red shirt is riding a wave
a man in a blue helmet is swimming pool on the ocean
two boys in goggles in a wave
100/100 [==============================] - 5s 55ms/step - loss: 3.1709 - masked_acc: 0.3830 - val_loss: 3.0774 - val_masked_acc: 0.3817
Epoch 28/100
100/100 [==============================] - ETA: 0s - loss: 3.1661 - masked_acc: 0.3816
a man in a red shirt is riding a wave
a man is being at a wave
a man wearing a blue shirt helmet is jumping into a pool
100/100 [==============================] - 5s 53ms/step - loss: 3.1661 - masked_acc: 0.3816 - val_loss: 3.0868 - val_masked_acc: 0.3829
Epoch 29/100
98/100 [============================>.] - ETA: 0s - loss: 3.0804 - masked_acc: 0.3906
a man in a red shirt is swimming in the water
a man in a yellow hat is surfing on a wave
a surfer in the air from the kayak
100/100 [==============================] - 6s 55ms/step - loss: 3.0812 - masked_acc: 0.3904 - val_loss: 3.0592 - val_masked_acc: 0.3798
Epoch 30/100
100/100 [==============================] - ETA: 0s - loss: 3.0275 - masked_acc: 0.3964
a man in a red shirt is swimming in the ocean
a man with a yellow jacket is riding a wave
an leather woman wearing orange dress looks down a branch direction and pier in the background
100/100 [==============================] - 6s 59ms/step - loss: 3.0275 - masked_acc: 0.3964 - val_loss: 3.1164 - val_masked_acc: 0.3791
Epoch 31/100
99/100 [============================>.] - ETA: 0s - loss: 3.0601 - masked_acc: 0.3887
a man in a red shirt is riding a wave
a man in a blue jacket is splashing in the ocean
a teenage girl takes a under a handstand in the water
100/100 [==============================] - 5s 54ms/step - loss: 3.0610 - masked_acc: 0.3887 - val_loss: 3.0967 - val_masked_acc: 0.3784
Epoch 32/100
99/100 [============================>.] - ETA: 0s - loss: 3.0615 - masked_acc: 0.3897
a man in a red shirt is riding a wave
a man is surfing on a wave
woman in a red shirt biking in brown pink swinging by a in a sleeping
100/100 [==============================] - 6s 56ms/step - loss: 3.0597 - masked_acc: 0.3898 - val_loss: 3.0367 - val_masked_acc: 0.3847
Epoch 33/100
99/100 [============================>.] - ETA: 0s - loss: 3.0125 - masked_acc: 0.3964
a man in a blue wetsuit is riding a wave
a man in a red wetsuit is being wave
a child in green swim trunks swims in the water
100/100 [==============================] - 5s 52ms/step - loss: 3.0143 - masked_acc: 0.3958 - val_loss: 3.0568 - val_masked_acc: 0.3852
Epoch 34/100
99/100 [============================>.] - ETA: 0s - loss: 2.9823 - masked_acc: 0.3984
a man in a red shirt is riding a wave
a man in a yellow shirt is riding a wave
a man and goggles on the wave
100/100 [==============================] - 5s 53ms/step - loss: 2.9825 - masked_acc: 0.3985 - val_loss: 3.0075 - val_masked_acc: 0.3875
Epoch 35/100
99/100 [============================>.] - ETA: 0s - loss: 3.0006 - masked_acc: 0.3967
a man in a red wetsuit is riding a wave
a surfer is riding a wave
a person thrown its santa and blue wetsuit in a wave
100/100 [==============================] - 5s 51ms/step - loss: 3.0028 - masked_acc: 0.3965 - val_loss: 3.0462 - val_masked_acc: 0.3846
Epoch 36/100
99/100 [============================>.] - ETA: 0s - loss: 2.9830 - masked_acc: 0.3976
a man in a yellow shirt is surfing
a man is falling into the water
an orange and a surfer is riding a wave
100/100 [==============================] - 5s 49ms/step - loss: 2.9798 - masked_acc: 0.3980 - val_loss: 3.0503 - val_masked_acc: 0.3793
Epoch 37/100
99/100 [============================>.] - ETA: 0s - loss: 2.9679 - masked_acc: 0.4021
a man in a yellow kayak is riding a wave
a man in a blue wetsuit is riding a wave
a child skier
100/100 [==============================] - 5s 49ms/step - loss: 2.9660 - masked_acc: 0.4024 - val_loss: 2.9918 - val_masked_acc: 0.3895
Epoch 38/100
99/100 [============================>.] - ETA: 0s - loss: 2.9519 - masked_acc: 0.3999
a man in a red wetsuit is surfing
two people in a surfboard and goggles
the two men in the air on the wave
100/100 [==============================] - 5s 51ms/step - loss: 2.9514 - masked_acc: 0.3998 - val_loss: 3.0695 - val_masked_acc: 0.3821
Epoch 39/100
99/100 [============================>.] - ETA: 0s - loss: 2.8850 - masked_acc: 0.4076
a man in a red wetsuit is surfing
a surfer is riding a wave
this person rides a stunt wave
100/100 [==============================] - 5s 47ms/step - loss: 2.8861 - masked_acc: 0.4074 - val_loss: 3.0209 - val_masked_acc: 0.3845
Epoch 40/100
100/100 [==============================] - ETA: 0s - loss: 2.8455 - masked_acc: 0.4110
a man in a yellow kayak is surfing
a man in a yellow kayak is surfing
surfer while an child stands in the ocean
100/100 [==============================] - 5s 50ms/step - loss: 2.8455 - masked_acc: 0.4110 - val_loss: 3.0512 - val_masked_acc: 0.3741
Epoch 41/100
99/100 [============================>.] - ETA: 0s - loss: 2.8658 - masked_acc: 0.4102
a man in a yellow kayak is riding a wave
a man in a blue surfboard is surfing
a surfer is surfing on a wave
100/100 [==============================] - 5s 51ms/step - loss: 2.8632 - masked_acc: 0.4106 - val_loss: 2.9706 - val_masked_acc: 0.3903
Epoch 42/100
100/100 [==============================] - ETA: 0s - loss: 2.8930 - masked_acc: 0.4046
a man in a yellow kayak is surfing
a man in a red surfboard is surfing in the ocean
the surfer in a yellow helmet is shooting wave
100/100 [==============================] - 5s 54ms/step - loss: 2.8930 - masked_acc: 0.4046 - val_loss: 2.9379 - val_masked_acc: 0.4046
Epoch 43/100
100/100 [==============================] - ETA: 0s - loss: 2.8188 - masked_acc: 0.4140
a man in a yellow kayak is riding a wave
a man in a red helmet is surfing
a white dog stands down a wave
100/100 [==============================] - 5s 53ms/step - loss: 2.8188 - masked_acc: 0.4140 - val_loss: 2.8918 - val_masked_acc: 0.3988
Epoch 44/100
100/100 [==============================] - ETA: 0s - loss: 2.8512 - masked_acc: 0.4084
a man in a yellow shirt is riding a wave
a man in a yellow kayak is doing a wave
a man jumps in the green is float on her neck
100/100 [==============================] - 5s 54ms/step - loss: 2.8512 - masked_acc: 0.4084 - val_loss: 2.9555 - val_masked_acc: 0.3926
Epoch 45/100
100/100 [==============================] - ETA: 0s - loss: 2.8523 - masked_acc: 0.4072
a man in a yellow kayak is surfing
a man in a blue wave as he is surfing
a person is surfing in a wave
100/100 [==============================] - 5s 50ms/step - loss: 2.8523 - masked_acc: 0.4072 - val_loss: 2.9054 - val_masked_acc: 0.3986
Epoch 46/100
99/100 [============================>.] - ETA: 0s - loss: 2.8613 - masked_acc: 0.4111
a man in a yellow kayak is surfing
a man in a wetsuit is riding a wave
a surfer makes against a wave
100/100 [==============================] - 5s 49ms/step - loss: 2.8594 - masked_acc: 0.4115 - val_loss: 2.9277 - val_masked_acc: 0.3935
Epoch 47/100
100/100 [==============================] - ETA: 0s - loss: 2.8498 - masked_acc: 0.4072
a man in a red wetsuit is riding a wave
a surfer in a red life jacket is surfing
a man in a wetsuit rides a wave
100/100 [==============================] - 5s 53ms/step - loss: 2.8498 - masked_acc: 0.4072 - val_loss: 2.9910 - val_masked_acc: 0.3846
Epoch 48/100
99/100 [============================>.] - ETA: 0s - loss: 2.7612 - masked_acc: 0.4158
a man in a yellow kayak is riding a wave
a person in a yellow life jacket on a surfboard
man for road kayaking in the ocean
100/100 [==============================] - 5s 53ms/step - loss: 2.7613 - masked_acc: 0.4158 - val_loss: 2.9676 - val_masked_acc: 0.3871
[[["이해하기 쉬움","easyToUnderstand","thumb-up"],["문제가 해결됨","solvedMyProblem","thumb-up"],["기타","otherUp","thumb-up"]],[["필요한 정보가 없음","missingTheInformationINeed","thumb-down"],["너무 복잡함/단계 수가 너무 많음","tooComplicatedTooManySteps","thumb-down"],["오래됨","outOfDate","thumb-down"],["번역 문제","translationIssue","thumb-down"],["샘플/코드 문제","samplesCodeIssue","thumb-down"],["기타","otherDown","thumb-down"]],["최종 업데이트: 2024-02-12(UTC)"],[],[]]